layernorm_kernels.cu 12.8 KB
Newer Older
1
#include <torch/all.h>
2
#include <ATen/cuda/CUDAContext.h>
3
#include <c10/cuda/CUDAGuard.h>
zhuwenwen's avatar
zhuwenwen committed
4

5
#include "dispatch_utils.h"
Woosuk Kwon's avatar
Woosuk Kwon committed
6
#include "reduction_utils.cuh"
7
8
9
10
11
12
13
#ifndef USE_ROCM
  #include <cuda_bf16.h>
  #include <cuda_fp16.h>
#else
  #include <hip/hip_bf16.h>
  #include <hip/hip_fp16.h>

14
15
using __nv_bfloat16 = __hip_bfloat16;
using __nv_bfloat162 = __hip_bfloat162;
16
#endif
17

Woosuk Kwon's avatar
Woosuk Kwon committed
18
namespace vllm {
19
20

// TODO(woosuk): Further optimize this kernel.
21
template <typename scalar_t>
22
__global__ void rms_norm_kernel(
23
24
25
26
    scalar_t* __restrict__ out,           // [..., hidden_size]
    const scalar_t* __restrict__ input,   // [..., hidden_size]
    const scalar_t* __restrict__ weight,  // [hidden_size]
    const float epsilon, const int num_tokens, const int hidden_size) {
27
28
29
30
  __shared__ float s_variance;
  float variance = 0.0f;

  for (int idx = threadIdx.x; idx < hidden_size; idx += blockDim.x) {
31
    const float x = (float)input[blockIdx.x * hidden_size + idx];
32
33
34
35
36
37
38
39
40
    variance += x * x;
  }
  variance = blockReduceSum<float>(variance);
  if (threadIdx.x == 0) {
    s_variance = rsqrtf(variance / hidden_size + epsilon);
  }
  __syncthreads();

  for (int idx = threadIdx.x; idx < hidden_size; idx += blockDim.x) {
41
42
43
    float x = (float)input[blockIdx.x * hidden_size + idx];
    out[blockIdx.x * hidden_size + idx] =
        ((scalar_t)(x * s_variance)) * weight[idx];
44
45
46
  }
}

47
48
49
50
51
52
53
54
/* Converter structs for the conversion from torch types to HIP/CUDA types,
   and the associated type conversions within HIP/CUDA. These helpers need
   to be implemented for now because the relevant type conversion
   operators/constructors are not consistently implemented by HIP/CUDA, so
   a generic conversion via type casts cannot be implemented.

   Each struct should have the member static constexpr bool `exists`:
   If false, the optimized kernel is not used for the corresponding torch type.
55
   If true, the struct should be fully defined as shown in the examples below.
56
 */
57
58
59
60
template <typename torch_type>
struct _typeConvert {
  static constexpr bool exists = false;
};
61

62
63
#if defined(USE_ROCM) || (defined(CUDA_VERSION) && (CUDA_VERSION >= 12000))
// CUDA < 12.0 runs into issues with packed type conversion
64
template <>
65
66
67
68
69
70
struct _typeConvert<c10::Half> {
  static constexpr bool exists = true;
  using hip_type = __half;
  using packed_hip_type = __half2;

  __device__ static inline float convert(hip_type x) { return __half2float(x); }
71
72
73
74
75
76
77
78
79
  __device__ static inline float2 convert(packed_hip_type x) {
    return __half22float2(x);
  }
  __device__ static inline hip_type convert(float x) {
    return __float2half_rn(x);
  }
  __device__ static inline packed_hip_type convert(float2 x) {
    return __float22half2_rn(x);
  }
80
81
};

82
  #if defined(__CUDA_ARCH__) && __CUDA_ARCH__ >= 800
83
84
// CUDA_ARCH < 800 does not have BF16 support
// TODO: Add in ROCm support once public headers handle bf16 maturely
85
template <>
86
87
88
89
90
struct _typeConvert<c10::BFloat16> {
  static constexpr bool exists = true;
  using hip_type = __nv_bfloat16;
  using packed_hip_type = __nv_bfloat162;

91
92
93
94
95
96
97
98
99
100
101
102
  __device__ static inline float convert(hip_type x) {
    return __bfloat162float(x);
  }
  __device__ static inline float2 convert(packed_hip_type x) {
    return __bfloat1622float2(x);
  }
  __device__ static inline hip_type convert(float x) {
    return __float2bfloat16(x);
  }
  __device__ static inline packed_hip_type convert(float2 x) {
    return __float22bfloat162_rn(x);
  }
103
};
104
105
106
  #endif  // defined(__CUDA_ARCH__) && __CUDA_ARCH__ >= 800
#endif    // defined(USE_ROCM) || (defined(CUDA_VERSION) && (CUDA_VERSION >=
          // 12000))
107
108
109
110
111
112

/* Vector POD struct to generate vectorized and packed FP16/BF16 ops
   for appropriate specializations of fused_add_rms_norm_kernel.
   Only functions that are necessary in that kernel are implemented.
   Alignment to 16 bytes is required to use 128-bit global memory ops.
 */
113
template <typename scalar_t, int width>
114
struct alignas(16) _f16Vec {
115
116
  /* Not theoretically necessary that width is a power of 2 but should
     almost always be the case for optimization purposes */
117
118
119
120
121
122
123
124
125
  static_assert(width > 0 && (width & (width - 1)) == 0,
                "Width is not a positive power of 2!");
  using Converter = _typeConvert<scalar_t>;
  using T1 = typename Converter::hip_type;
  using T2 = typename Converter::packed_hip_type;
  T1 data[width];

  __device__ _f16Vec& operator+=(const _f16Vec<scalar_t, width>& other) {
    if constexpr (width % 2 == 0) {
126
#pragma unroll
127
      for (int i = 0; i < width; i += 2) {
128
129
        T2 temp{data[i], data[i + 1]};
        temp += T2{other.data[i], other.data[i + 1]};
130
        data[i] = temp.x;
131
        data[i + 1] = temp.y;
132
133
      }
    } else {
134
135
#pragma unroll
      for (int i = 0; i < width; ++i) data[i] += other.data[i];
136
137
138
139
140
141
    }
    return *this;
  }

  __device__ _f16Vec& operator*=(const _f16Vec<scalar_t, width>& other) {
    if constexpr (width % 2 == 0) {
142
#pragma unroll
143
      for (int i = 0; i < width; i += 2) {
144
145
        T2 temp{data[i], data[i + 1]};
        temp *= T2{other.data[i], other.data[i + 1]};
146
        data[i] = temp.x;
147
        data[i + 1] = temp.y;
148
149
      }
    } else {
150
151
#pragma unroll
      for (int i = 0; i < width; ++i) data[i] *= other.data[i];
152
153
154
155
156
157
    }
    return *this;
  }

  __device__ _f16Vec& operator*=(const float scale) {
    if constexpr (width % 2 == 0) {
158
#pragma unroll
159
      for (int i = 0; i < width; i += 2) {
160
        float2 temp_f = Converter::convert(T2{data[i], data[i + 1]});
161
162
163
164
        temp_f.x *= scale;
        temp_f.y *= scale;
        T2 temp = Converter::convert(temp_f);
        data[i] = temp.x;
165
        data[i + 1] = temp.y;
166
167
      }
    } else {
168
#pragma unroll
169
170
171
172
173
174
175
176
177
178
179
      for (int i = 0; i < width; ++i) {
        float temp = Converter::convert(data[i]) * scale;
        data[i] = Converter::convert(temp);
      }
    }
    return *this;
  }

  __device__ float sum_squares() const {
    float result = 0.0f;
    if constexpr (width % 2 == 0) {
180
#pragma unroll
181
      for (int i = 0; i < width; i += 2) {
182
        float2 z = Converter::convert(T2{data[i], data[i + 1]});
183
184
185
        result += z.x * z.x + z.y * z.y;
      }
    } else {
186
#pragma unroll
187
188
189
190
191
192
193
194
195
196
197
198
199
      for (int i = 0; i < width; ++i) {
        float x = Converter::convert(data[i]);
        result += x * x;
      }
    }
    return result;
  }
};

/* Function specialization in the case of FP16/BF16 tensors.
   Additional optimizations we can make in this case are
   packed and vectorized operations, which help with the
   memory latency bottleneck. */
200
201
202
203
204
205
206
template <typename scalar_t, int width>
__global__ std::enable_if_t<(width > 0) && _typeConvert<scalar_t>::exists>
fused_add_rms_norm_kernel(
    scalar_t* __restrict__ input,         // [..., hidden_size]
    scalar_t* __restrict__ residual,      // [..., hidden_size]
    const scalar_t* __restrict__ weight,  // [hidden_size]
    const float epsilon, const int num_tokens, const int hidden_size) {
207
208
209
210
211
212
213
214
215
216
  // Sanity checks on our vector struct and type-punned pointer arithmetic
  static_assert(std::is_pod_v<_f16Vec<scalar_t, width>>);
  static_assert(sizeof(_f16Vec<scalar_t, width>) == sizeof(scalar_t) * width);

  const int vec_hidden_size = hidden_size / width;
  __shared__ float s_variance;
  float variance = 0.0f;
  /* These and the argument pointers are all declared `restrict` as they are
     not aliased in practice. Argument pointers should not be dereferenced
     in this kernel as that would be undefined behavior */
217
218
219
220
221
222
  auto* __restrict__ input_v =
      reinterpret_cast<_f16Vec<scalar_t, width>*>(input);
  auto* __restrict__ residual_v =
      reinterpret_cast<_f16Vec<scalar_t, width>*>(residual);
  auto* __restrict__ weight_v =
      reinterpret_cast<const _f16Vec<scalar_t, width>*>(weight);
223
224
225
226
227
228
229
230
231

  for (int idx = threadIdx.x; idx < vec_hidden_size; idx += blockDim.x) {
    int id = blockIdx.x * vec_hidden_size + idx;
    _f16Vec<scalar_t, width> temp = input_v[id];
    temp += residual_v[id];
    variance += temp.sum_squares();
    residual_v[id] = temp;
  }
  /* Keep the following if-else block in sync with the
232
     calculation of max_block_size in fused_add_rms_norm */
233
234
  if (num_tokens < 256) {
    variance = blockReduceSum<float, 1024>(variance);
235
236
  } else
    variance = blockReduceSum<float, 256>(variance);
237
238
239
240
241
242
243
244
245
246
247
248
249
250
251
252
253
  if (threadIdx.x == 0) {
    s_variance = rsqrtf(variance / hidden_size + epsilon);
  }
  __syncthreads();

  for (int idx = threadIdx.x; idx < vec_hidden_size; idx += blockDim.x) {
    int id = blockIdx.x * vec_hidden_size + idx;
    _f16Vec<scalar_t, width> temp = residual_v[id];
    temp *= s_variance;
    temp *= weight_v[idx];
    input_v[id] = temp;
  }
}

/* Generic fused_add_rms_norm_kernel
   The width field is not used here but necessary for other specializations.
 */
254
255
256
257
258
259
260
template <typename scalar_t, int width>
__global__ std::enable_if_t<(width == 0) || !_typeConvert<scalar_t>::exists>
fused_add_rms_norm_kernel(
    scalar_t* __restrict__ input,         // [..., hidden_size]
    scalar_t* __restrict__ residual,      // [..., hidden_size]
    const scalar_t* __restrict__ weight,  // [hidden_size]
    const float epsilon, const int num_tokens, const int hidden_size) {
261
262
263
264
  __shared__ float s_variance;
  float variance = 0.0f;

  for (int idx = threadIdx.x; idx < hidden_size; idx += blockDim.x) {
265
266
    scalar_t z = input[blockIdx.x * hidden_size + idx];
    z += residual[blockIdx.x * hidden_size + idx];
267
    float x = (float)z;
268
    variance += x * x;
269
    residual[blockIdx.x * hidden_size + idx] = z;
270
  }
271
  /* Keep the following if-else block in sync with the
272
     calculation of max_block_size in fused_add_rms_norm */
273
274
  if (num_tokens < 256) {
    variance = blockReduceSum<float, 1024>(variance);
275
276
  } else
    variance = blockReduceSum<float, 256>(variance);
277
278
279
280
281
282
  if (threadIdx.x == 0) {
    s_variance = rsqrtf(variance / hidden_size + epsilon);
  }
  __syncthreads();

  for (int idx = threadIdx.x; idx < hidden_size; idx += blockDim.x) {
283
284
285
    float x = (float)residual[blockIdx.x * hidden_size + idx];
    input[blockIdx.x * hidden_size + idx] =
        ((scalar_t)(x * s_variance)) * weight[idx];
286
287
288
  }
}

289
}  // namespace vllm
290

291
292
293
void rms_norm(torch::Tensor& out,     // [..., hidden_size]
              torch::Tensor& input,   // [..., hidden_size]
              torch::Tensor& weight,  // [hidden_size]
294
              double epsilon) {
295
296
  int hidden_size = input.size(-1);
  int num_tokens = input.numel() / hidden_size;
zhuwenwen's avatar
zhuwenwen committed
297
298
299

  dim3 grid(num_tokens);
  dim3 block(std::min(hidden_size, 1024));
300
  const at::cuda::OptionalCUDAGuard device_guard(device_of(input));
301
  const cudaStream_t stream = at::cuda::getCurrentCUDAStream();
zhuwenwen's avatar
zhuwenwen committed
302
303
304
305
306
  VLLM_DISPATCH_FLOATING_TYPES(input.scalar_type(), "rms_norm_kernel", [&] {
    vllm::rms_norm_kernel<scalar_t><<<grid, block, 0, stream>>>(
        out.data_ptr<scalar_t>(), input.data_ptr<scalar_t>(),
        weight.data_ptr<scalar_t>(), epsilon, num_tokens, hidden_size);
  });
307
}
308

309
310
311
312
313
314
315
316
317
318
319
320
321
#define LAUNCH_FUSED_ADD_RMS_NORM(width)                                       \
  VLLM_DISPATCH_FLOATING_TYPES(                                                \
      input.scalar_type(), "fused_add_rms_norm_kernel", [&] {                  \
        vllm::fused_add_rms_norm_kernel<scalar_t, width>                       \
            <<<grid, block, 0, stream>>>(input.data_ptr<scalar_t>(),           \
                                         residual.data_ptr<scalar_t>(),        \
                                         weight.data_ptr<scalar_t>(), epsilon, \
                                         num_tokens, hidden_size);             \
      });

void fused_add_rms_norm(torch::Tensor& input,     // [..., hidden_size]
                        torch::Tensor& residual,  // [..., hidden_size]
                        torch::Tensor& weight,    // [hidden_size]
322
                        double epsilon) {
323
324
  int hidden_size = input.size(-1);
  int num_tokens = input.numel() / hidden_size;
zhuwenwen's avatar
zhuwenwen committed
325
326
327
328
329
330
331
332

  dim3 grid(num_tokens);
  /* This kernel is memory-latency bound in many scenarios.
     When num_tokens is large, a smaller block size allows
     for increased block occupancy on CUs and better latency
     hiding on global mem ops. */
  const int max_block_size = (num_tokens < 256) ? 1024 : 256;
  dim3 block(std::min(hidden_size, max_block_size));
333
  const at::cuda::OptionalCUDAGuard device_guard(device_of(input));
334
  const cudaStream_t stream = at::cuda::getCurrentCUDAStream();
zhuwenwen's avatar
zhuwenwen committed
335
336
337
338
339
340
341
342
343
344
345
346
347
348
349
350
  /*If the tensor types are FP16/BF16, try to use the optimized kernel
    with packed + vectorized ops.
    Max optimization is achieved with a width-8 vector of FP16/BF16s
    since we can load at most 128 bits at once in a global memory op.
    However, this requires each tensor's data to be aligned to 16
    bytes.
   */
  auto inp_ptr = reinterpret_cast<std::uintptr_t>(input.data_ptr());
  auto res_ptr = reinterpret_cast<std::uintptr_t>(residual.data_ptr());
  auto wt_ptr = reinterpret_cast<std::uintptr_t>(weight.data_ptr());
  bool ptrs_are_aligned =
      inp_ptr % 16 == 0 && res_ptr % 16 == 0 && wt_ptr % 16 == 0;
  if (ptrs_are_aligned && hidden_size % 8 == 0) {
    LAUNCH_FUSED_ADD_RMS_NORM(8);
  } else {
    LAUNCH_FUSED_ADD_RMS_NORM(0);
zhangshao's avatar
zhangshao committed
351
  }
zhuwenwen's avatar
zhuwenwen committed
352
}