File size: 19,078 Bytes
d90b3a8
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
1
2
3
4
5
6
7
8
9
10
11
12
13
14
15
16
17
18
19
20
21
22
23
24
25
26
27
28
29
30
31
32
33
34
35
36
37
38
39
40
41
42
43
44
45
46
47
48
49
50
51
52
53
54
55
56
57
58
59
60
61
62
63
64
65
66
67
68
69
70
71
72
73
74
75
76
77
78
79
80
81
82
83
84
85
86
87
88
89
90
91
92
93
94
95
96
97
98
99
100
101
102
103
104
105
106
107
108
109
110
111
112
113
114
115
116
117
118
119
120
121
122
123
124
125
126
127
128
129
130
131
132
133
134
135
136
137
138
139
140
141
142
143
144
145
146
147
148
149
150
151
152
153
154
155
156
157
158
159
160
161
162
163
164
165
166
167
168
169
170
171
172
173
174
175
176
177
178
179
180
181
182
183
184
185
186
187
188
189
190
191
192
193
194
195
196
197
198
199
200
201
202
203
204
205
206
207
208
209
210
211
212
213
214
215
216
217
218
219
220
221
222
223
224
225
226
227
228
229
230
231
232
233
234
235
236
237
238
239
240
241
242
243
244
245
246
247
248
249
250
251
252
253
254
255
256
257
258
259
260
261
262
263
264
265
266
267
268
269
270
271
272
273
274
275
276
277
278
279
280
281
282
283
284
285
286
287
288
289
290
291
292
293
294
295
296
297
298
299
300
301
302
303
304
305
306
307
308
309
310
311
312
313
314
315
316
317
318
319
320
321
322
323
324
325
326
327
328
329
330
331
332
333
334
335
336
337
338
339
340
341
342
343
344
345
346
347
348
349
350
351
352
353
354
355
356
357
358
359
360
361
362
363
364
365
366
367
368
369
370
371
372
373
374
375
376
377
378
379
380
381
382
383
384
385
386
387
388
389
390
391
392
393
394
395
396
/* coding=utf-8
 * Copyright (c) 2024, NVIDIA CORPORATION.  All rights reserved.
 *
 * Licensed under the Apache License, Version 2.0 (the "License");
 * you may not use this file except in compliance with the License.
 * You may obtain a copy of the License at
 *
 *     http://www.apache.org/licenses/LICENSE-2.0
 *
 * Unless required by applicable law or agreed to in writing, software
 * distributed under the License is distributed on an "AS IS" BASIS,
 * WITHOUT WARRANTIES OR CONDITIONS OF ANY KIND, either express or implied.
 * See the License for the specific language governing permissions and
 * limitations under the License.
 */

#pragma once

#include <ATen/ATen.h>
#include <ATen/cuda/CUDAContext.h>
#include <c10/macros/Macros.h>
#include <cuda_runtime.h>
#include <torch/extension.h>

namespace {

template <typename scalar_t>
__global__ void fused_rope_forward(const int h,
                                   const int d,
                                   const int d2,
                                   const int stride_s,
                                   const int stride_b,
                                   const int stride_h,
                                   const int stride_d,
                                   const int o_stride_s,
                                   const int o_stride_b,
                                   const int o_stride_h,
                                   const int o_stride_d,
                                   const scalar_t* src,
                                   const float* freqs,
                                   scalar_t* dst)
{
    int s_id = blockIdx.x, b_id = blockIdx.y;
    int offset_block = s_id * stride_s + b_id * stride_b;
    int offset_block_dst = s_id * o_stride_s + b_id * o_stride_b;
#pragma unroll
    for (int d_id = threadIdx.x; d_id < d2; d_id += blockDim.x) {
        float v_cos, v_sin;
        sincosf(freqs[s_id * d2 + d_id], &v_sin, &v_cos);
#pragma unroll
        for (int h_id = threadIdx.y; h_id < h; h_id += blockDim.y) {
            int offset_src = offset_block + h_id * stride_h + d_id * stride_d;
            int offset_dst = offset_block_dst + h_id * o_stride_h + d_id * o_stride_d;
            scalar_t v_src = src[offset_src];
            scalar_t v_src_rotate = (d_id + d2 / 2 < d2)
                                        ? -src[offset_src + (d2 / 2) * stride_d]
                                        : src[offset_src + (d2 / 2 - d2) * stride_d];
            dst[offset_dst] = v_src * (scalar_t)v_cos + v_src_rotate * (scalar_t)v_sin;
        }
    }

    // copy the rest
    if (d > d2) {
#pragma unroll
        for (int h_id = threadIdx.y; h_id < h; h_id += blockDim.y) {
            int offset_head = offset_block + h_id * stride_h;
            int offset_head_dst = offset_block_dst + h_id * o_stride_h;
#pragma unroll
            for (int d_id = d2 + threadIdx.x; d_id < d; d_id += blockDim.x) {
                dst[offset_head_dst + d_id * o_stride_d] = src[offset_head + d_id * stride_d];
            }
        }
    }
}

template <typename scalar_t>
__global__ void fused_rope_backward(const int h,
                                    const int d,
                                    const int d2,
                                    const int stride_s,
                                    const int stride_b,
                                    const int stride_h,
                                    const int stride_d,
                                    const int o_stride_s,
                                    const int o_stride_b,
                                    const int o_stride_h,
                                    const int o_stride_d,
                                    const scalar_t* src,
                                    const float* freqs,
                                    scalar_t* dst)
{
    int s_id = blockIdx.x, b_id = blockIdx.y;
    int offset_block = s_id * stride_s + b_id * stride_b;
    int offset_block_dst = s_id * o_stride_s + b_id * o_stride_b;
#pragma unroll
    for (int d_id = threadIdx.x; d_id < d2; d_id += blockDim.x) {
        scalar_t v_cos = cosf(freqs[s_id * d2 + d_id]);
        scalar_t v_sin = (d_id + d2 / 2 < d2) ? sinf(freqs[s_id * d2 + d_id + d2 / 2])
                                              : -sinf(freqs[s_id * d2 + d_id + d2 / 2 - d2]);
#pragma unroll
        for (int h_id = threadIdx.y; h_id < h; h_id += blockDim.y) {
            int offset_src = offset_block + h_id * stride_h + d_id * stride_d;
            int offset_dst = offset_block_dst + h_id * o_stride_h + d_id * o_stride_d;
            scalar_t v_src = src[offset_src];
            scalar_t v_src_rotate = (d_id + d2 / 2 < d2)
                                        ? src[offset_src + (d2 / 2) * stride_d]
                                        : src[offset_src + (d2 / 2 - d2) * stride_d];
            dst[offset_dst] = v_src * v_cos + v_src_rotate * v_sin;
        }
    }

    // handle the tail
    if (d > d2) {
#pragma unroll
        for (int h_id = threadIdx.y; h_id < h; h_id += blockDim.y) {
            int offset_head = offset_block + h_id * stride_h;
            int offset_head_dst = offset_block_dst + h_id * o_stride_h;
#pragma unroll
            for (int d_id = d2 + threadIdx.x; d_id < d; d_id += blockDim.x) {
                dst[offset_head_dst + d_id * o_stride_d] = src[offset_head + d_id * stride_d];
            }
        }
    }
}

template <typename scalar_t_0, typename scalar_t_1>
__global__ void fused_rope_cached_forward(const int h,
                                          const int d,
                                          const int d2,
                                          const int stride_s,
                                          const int stride_b,
                                          const int stride_h,
                                          const int stride_d,
                                          const int o_stride_s,
                                          const int o_stride_b,
                                          const int o_stride_h,
                                          const int o_stride_d,
                                          const scalar_t_0* src,
                                          const scalar_t_1* cos,
                                          const scalar_t_1* sin,
                                          scalar_t_0* dst)
{
    int s_id = blockIdx.x, b_id = blockIdx.y;
    int offset_block = s_id * stride_s + b_id * stride_b;
    int offset_block_dst = s_id * o_stride_s + b_id * o_stride_b;
#pragma unroll
    for (int d_id = threadIdx.x; d_id < d2; d_id += blockDim.x) {
        scalar_t_0 v_cos = cos[s_id * d2 + d_id];
        scalar_t_0 v_sin = sin[s_id * d2 + d_id];
#pragma unroll
        for (int h_id = threadIdx.y; h_id < h; h_id += blockDim.y) {
            int offset_src = offset_block + h_id * stride_h + d_id * stride_d;
            int offset_dst = offset_block_dst + h_id * o_stride_h + d_id * o_stride_d;
            scalar_t_0 v_src = src[offset_src];
            scalar_t_0 v_src_rotate = (d_id + d2 / 2 < d2)
                                          ? -src[offset_src + (d2 / 2) * stride_d]
                                          : src[offset_src + (d2 / 2 - d2) * stride_d];
            dst[offset_dst] = v_src * v_cos + v_src_rotate * v_sin;
        }
    }

    // copy the rest
    if (d > d2) {
#pragma unroll
        for (int h_id = threadIdx.y; h_id < h; h_id += blockDim.y) {
            int offset_head = offset_block + h_id * stride_h;
            int offset_head_dst = offset_block_dst + h_id * o_stride_h;
#pragma unroll
            for (int d_id = d2 + threadIdx.x; d_id < d; d_id += blockDim.x) {
                dst[offset_head_dst + d_id * o_stride_d] = src[offset_head + d_id * stride_d];
            }
        }
    }
}

template <typename scalar_t_0, typename scalar_t_1>
__global__ void fused_rope_cached_backward(const int h,
                                           const int d,
                                           const int d2,
                                           const int stride_s,
                                           const int stride_b,
                                           const int stride_h,
                                           const int stride_d,
                                           const int o_stride_s,
                                           const int o_stride_b,
                                           const int o_stride_h,
                                           const int o_stride_d,
                                           const scalar_t_0* src,
                                           const scalar_t_1* cos,
                                           const scalar_t_1* sin,
                                           scalar_t_0* dst)
{
    int s_id = blockIdx.x, b_id = blockIdx.y;
    int offset_block = s_id * stride_s + b_id * stride_b;
    int offset_block_dst = s_id * o_stride_s + b_id * o_stride_b;
#pragma unroll
    for (int d_id = threadIdx.x; d_id < d2; d_id += blockDim.x) {
        scalar_t_0 v_cos = cos[s_id * d2 + d_id];
        scalar_t_0 v_sin = (d_id + d2 / 2 < d2) ? sin[s_id * d2 + d_id + d2 / 2]
                                                : -sin[s_id * d2 + d_id + d2 / 2 - d2];
#pragma unroll
        for (int h_id = threadIdx.y; h_id < h; h_id += blockDim.y) {
            int offset_src = offset_block + h_id * stride_h + d_id * stride_d;
            int offset_dst = offset_block_dst + h_id * o_stride_h + d_id * o_stride_d;
            scalar_t_0 v_src = src[offset_src];
            scalar_t_0 v_src_rotate = (d_id + d2 / 2 < d2)
                                          ? src[offset_src + (d2 / 2) * stride_d]
                                          : src[offset_src + (d2 / 2 - d2) * stride_d];
            dst[offset_dst] = v_src * v_cos + v_src_rotate * v_sin;
        }
    }

    // handle the tail
    if (d > d2) {
#pragma unroll
        for (int h_id = threadIdx.y; h_id < h; h_id += blockDim.y) {
            int offset_head = offset_block + h_id * stride_h;
            int offset_head_dst = offset_block_dst + h_id * o_stride_h;
#pragma unroll
            for (int d_id = d2 + threadIdx.x; d_id < d; d_id += blockDim.x) {
                dst[offset_head_dst + d_id * o_stride_d] = src[offset_head + d_id * stride_d];
            }
        }
    }
}

}  // end of anonymous namespace

template <typename scalar_t>
void dispatch_fused_rope_forward(const int s,
                                 const int b,
                                 const int h,
                                 const int d,
                                 const int d2,
                                 const int stride_s,
                                 const int stride_b,
                                 const int stride_h,
                                 const int stride_d,
                                 const int o_stride_s,
                                 const int o_stride_b,
                                 const int o_stride_h,
                                 const int o_stride_d,
                                 const scalar_t* input,
                                 const float* freqs,
                                 scalar_t* output)
{
    auto stream = at::cuda::getCurrentCUDAStream();

    int warps_per_block = h < 16 ? 4 : 8;
    dim3 blocks(s, b);
    dim3 threads(C10_WARP_SIZE, warps_per_block);

    fused_rope_forward<<<blocks, threads, 0, stream>>>(h,
                                                       d,
                                                       d2,
                                                       stride_s,
                                                       stride_b,
                                                       stride_h,
                                                       stride_d,
                                                       o_stride_s,
                                                       o_stride_b,
                                                       o_stride_h,
                                                       o_stride_d,
                                                       input,
                                                       freqs,
                                                       output);
    C10_CUDA_KERNEL_LAUNCH_CHECK();
}

template <typename scalar_t>
void dispatch_fused_rope_backward(const int s,
                                  const int b,
                                  const int h,
                                  const int d,
                                  const int d2,
                                  const int stride_s,
                                  const int stride_b,
                                  const int stride_h,
                                  const int stride_d,
                                  const int o_stride_s,
                                  const int o_stride_b,
                                  const int o_stride_h,
                                  const int o_stride_d,
                                  const scalar_t* output_grads,
                                  const float* freqs,
                                  scalar_t* input_grads)
{
    auto stream = at::cuda::getCurrentCUDAStream();

    int warps_per_block = h < 16 ? 4 : 8;
    dim3 blocks(s, b);
    dim3 threads(C10_WARP_SIZE, warps_per_block);

    fused_rope_backward<<<blocks, threads, 0, stream>>>(h,
                                                        d,
                                                        d2,
                                                        stride_s,
                                                        stride_b,
                                                        stride_h,
                                                        stride_d,
                                                        o_stride_s,
                                                        o_stride_b,
                                                        o_stride_h,
                                                        o_stride_d,
                                                        output_grads,
                                                        freqs,
                                                        input_grads);
    C10_CUDA_KERNEL_LAUNCH_CHECK();
}

template <typename scalar_t_0, typename scalar_t_1>
void dispatch_fused_rope_cached_forward(const int s,
                                        const int b,
                                        const int h,
                                        const int d,
                                        const int d2,
                                        const int stride_s,
                                        const int stride_b,
                                        const int stride_h,
                                        const int stride_d,
                                        const int o_stride_s,
                                        const int o_stride_b,
                                        const int o_stride_h,
                                        const int o_stride_d,
                                        const scalar_t_0* input,
                                        const scalar_t_1* cos,
                                        const scalar_t_1* sin,
                                        scalar_t_0* output)
{
    auto stream = at::cuda::getCurrentCUDAStream();

    int warps_per_block = h < 16 ? 4 : 8;
    dim3 blocks(s, b);
    dim3 threads(C10_WARP_SIZE, warps_per_block);

    fused_rope_cached_forward<<<blocks, threads, 0, stream>>>(h,
                                                              d,
                                                              d2,
                                                              stride_s,
                                                              stride_b,
                                                              stride_h,
                                                              stride_d,
                                                              o_stride_s,
                                                              o_stride_b,
                                                              o_stride_h,
                                                              o_stride_d,
                                                              input,
                                                              cos,
                                                              sin,
                                                              output);
    C10_CUDA_KERNEL_LAUNCH_CHECK();
}

template <typename scalar_t_0, typename scalar_t_1>
void dispatch_fused_rope_cached_backward(const int s,
                                         const int b,
                                         const int h,
                                         const int d,
                                         const int d2,
                                         const int stride_s,
                                         const int stride_b,
                                         const int stride_h,
                                         const int stride_d,
                                         const int o_stride_s,
                                         const int o_stride_b,
                                         const int o_stride_h,
                                         const int o_stride_d,
                                         const scalar_t_0* output_grads,
                                         const scalar_t_1* cos,
                                         const scalar_t_1* sin,
                                         scalar_t_0* input_grads)
{
    auto stream = at::cuda::getCurrentCUDAStream();

    int warps_per_block = h < 16 ? 4 : 8;
    dim3 blocks(s, b);
    dim3 threads(C10_WARP_SIZE, warps_per_block);

    fused_rope_cached_backward<<<blocks, threads, 0, stream>>>(h,
                                                               d,
                                                               d2,
                                                               stride_s,
                                                               stride_b,
                                                               stride_h,
                                                               stride_d,
                                                               o_stride_s,
                                                               o_stride_b,
                                                               o_stride_h,
                                                               o_stride_d,
                                                               output_grads,
                                                               cos,
                                                               sin,
                                                               input_grads);
    C10_CUDA_KERNEL_LAUNCH_CHECK();
}