"docs/understand_mmcv/io.md" did not exist on "f573c11c3db7d38baef4f780eb85cad4155461f8"
pos_encoding_kernels.cu 2.75 KB
Newer Older
1
2
3
4
5
6
7
8
#include <torch/extension.h>
#include <ATen/cuda/CUDAContext.h>

namespace cacheflow {

template<typename scalar_t>
__global__ void rotary_embedding_neox_kernel(
  const int64_t* __restrict__ positions,        // [num_tokens]
Woosuk Kwon's avatar
Woosuk Kwon committed
9
10
  scalar_t* __restrict__ query,                 // [num_tokens, num_heads, head_size]
  scalar_t* __restrict__ key,                   // [num_tokens, num_heads, head_size]
11
  const scalar_t* __restrict__ cos_sin_cache,   // [max_position, 2, head_size // 2]
Woosuk Kwon's avatar
Woosuk Kwon committed
12
  const int stride,
13
14
15
16
17
18
19
20
  const int num_heads,
  const int head_size) {
  // Each thread block is responsible for one token.
  const int token_idx = blockIdx.x;
  int64_t pos = positions[token_idx];
  const scalar_t* cache_ptr = cos_sin_cache + pos * head_size;

  const int embed_dim = head_size / 2;
Woosuk Kwon's avatar
Woosuk Kwon committed
21
  const int n = num_heads * embed_dim;
22
  for (int i = threadIdx.x; i < n; i += blockDim.x) {
Woosuk Kwon's avatar
Woosuk Kwon committed
23
24
    const int head_idx = i / embed_dim;
    const int token_head = token_idx * stride + head_idx * head_size;
25

Woosuk Kwon's avatar
Woosuk Kwon committed
26
    const int rot_offset = i % embed_dim;
27
28
29
    const int x_index = rot_offset;
    const int y_index = embed_dim + rot_offset;

Woosuk Kwon's avatar
Woosuk Kwon committed
30
31
32
    const int out_x = token_idx * stride + head_idx * head_size + x_index;
    const int out_y = token_idx * stride + head_idx * head_size + y_index;

33
34
35
    const scalar_t cos = __ldg(cache_ptr + x_index);
    const scalar_t sin = __ldg(cache_ptr + y_index);

Woosuk Kwon's avatar
Woosuk Kwon committed
36
37
38
39
    const scalar_t q_x = query[token_head + x_index];
    const scalar_t q_y = query[token_head + y_index];
    query[out_x] = q_x * cos - q_y * sin;
    query[out_y] = q_y * cos + q_x * sin;
40

Woosuk Kwon's avatar
Woosuk Kwon committed
41
42
43
44
    const scalar_t k_x = key[token_head + x_index];
    const scalar_t k_y = key[token_head + y_index];
    key[out_x] = k_x * cos - k_y * sin;
    key[out_y] = k_y * cos + k_x * sin;
45
46
47
48
49
50
51
52
53
54
55
56
57
58
  }
}

} // namespace cacheflow

void rotary_embedding_neox(
  torch::Tensor& positions,         // [num_tokens]
  torch::Tensor& query,             // [num_tokens, num_heads * head_size]
  torch::Tensor& key,               // [num_tokens, num_heads * head_size]
  torch::Tensor& cos_sin_cache)     // [max_position, head_size]
{
  int num_tokens = query.size(0);
  int head_size = cos_sin_cache.size(1);
  int num_heads = query.size(1) / head_size;
Woosuk Kwon's avatar
Woosuk Kwon committed
59
60
  int stride = query.stride(0);
  TORCH_CHECK(stride == key.stride(0));
61
62

  dim3 grid(num_tokens);
Woosuk Kwon's avatar
Woosuk Kwon committed
63
  dim3 block(std::min(num_heads * head_size / 2, 512));
64
65
66
67
68
69
70
71
72
73
  const cudaStream_t stream = at::cuda::getCurrentCUDAStream();
  AT_DISPATCH_FLOATING_TYPES_AND_HALF(
    query.scalar_type(),
    "rotary_embedding_neox",
    [&] {
      cacheflow::rotary_embedding_neox_kernel<scalar_t><<<grid, block, 0, stream>>>(
        positions.data_ptr<int64_t>(),
        query.data_ptr<scalar_t>(),
        key.data_ptr<scalar_t>(),
        cos_sin_cache.data_ptr<scalar_t>(),
Woosuk Kwon's avatar
Woosuk Kwon committed
74
        stride,
75
76
77
78
        num_heads,
        head_size);
    });
}