scaled_masked_softmax.h 21 KB

123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121122123124125126127128129130131132133134135136137138139140141142143144145146147148149150151152153154155156157158159160161162163164165166167168169170171172173174175176177178179180181182183184185186187188189190191192193194195196197198199200201202203204205206207208209210211212213214215216217218219220221222223224225226227228229230231232233234235236237238239240241242243244245246247248249250251252253254255256257258259260261262263264265266267268269270271272273274275276277278279280281282283284285286287288289290291292293294295296297298299300301302303304305306307308309310311312313314315316317318319320321322323324325326327328329330331332333334335336337338339340341342343344345346347348349350351352353354355356357358359360361362363364365366367368369370371372373374375376377378379380381382383384385386387388389390391392393394395396397398399400401402403404405406407408409410411412413414415416417418419420421422423424425426427428429430431432433434435436437438439440441442443444445446447448449450451452453454455456457458459460461462463464465466467468469470471472473474475476477478479480481482483484485486487488489490491492493
  1. /* coding=utf-8
  2. * Copyright (c) 2020, NVIDIA CORPORATION. All rights reserved.
  3. *
  4. * Licensed under the Apache License, Version 2.0 (the "License");
  5. * you may not use this file except in compliance with the License.
  6. * You may obtain a copy of the License at
  7. *
  8. * http://www.apache.org/licenses/LICENSE-2.0
  9. *
  10. * Unless required by applicable law or agreed to in writing, software
  11. * distributed under the License is distributed on an "AS IS" BASIS,
  12. * WITHOUT WARRANTIES OR CONDITIONS OF ANY KIND, either express or implied.
  13. * See the License for the specific language governing permissions and
  14. * limitations under the License.
  15. */
  16. #pragma once
  17. #include <assert.h>
  18. #include <cuda_fp16.h>
  19. #include <cfloat>
  20. #include <limits>
  21. #include <stdint.h>
  22. #include <cuda_fp16.h>
  23. #include <c10/macros/Macros.h>
  24. namespace {
  25. template <typename Datatype, int ELEMENTS_PER_LDG>
  26. __device__ __inline__ void copy_vector(Datatype *dst, const Datatype *src);
  27. template <>
  28. __device__ __inline__ void copy_vector<c10::BFloat16, 1>(c10::BFloat16 *dst, const c10::BFloat16 *src) { *dst = *src; }
  29. template <>
  30. __device__ __inline__ void copy_vector<c10::BFloat16, 4>(c10::BFloat16 *dst, const c10::BFloat16 *src) { *((float2*) dst) = *((float2*) src); }
  31. template <>
  32. __device__ __inline__ void copy_vector<c10::Half, 1>(c10::Half *dst, const c10::Half *src) { *dst = *src; }
  33. template <>
  34. __device__ __inline__ void copy_vector<c10::Half, 4>(c10::Half *dst, const c10::Half *src) { *((float2*) dst) = *((float2*) src); }
  35. template <>
  36. __device__ __inline__ void copy_vector<uint8_t, 1>(uint8_t *dst, const uint8_t *src) { *dst = *src; }
  37. template <>
  38. __device__ __inline__ void copy_vector<uint8_t, 4>(uint8_t *dst, const uint8_t *src) {*((half2*) dst) = *((half2*) src); }
  39. int log2_ceil(int value) {
  40. int log2_value = 0;
  41. while ((1 << log2_value) < value) ++log2_value;
  42. return log2_value;
  43. }
  44. template<typename T>
  45. struct Add {
  46. __device__ __forceinline__ T operator()(T a, T b) const {
  47. return a + b;
  48. }
  49. };
  50. template<typename T>
  51. struct Max {
  52. __device__ __forceinline__ T operator()(T a, T b) const {
  53. return a < b ? b : a;
  54. }
  55. };
  56. template <typename T>
  57. __device__ __forceinline__ T WARP_SHFL_XOR_NATIVE(T value, int laneMask, int width = warpSize, unsigned int mask = 0xffffffff)
  58. {
  59. #if CUDA_VERSION >= 9000
  60. return __shfl_xor_sync(mask, value, laneMask, width);
  61. #else
  62. return __shfl_xor(value, laneMask, width);
  63. #endif
  64. }
  65. template <typename acc_t, int WARP_BATCH, int WARP_SIZE, template<typename> class ReduceOp>
  66. __device__ __forceinline__ void warp_reduce(acc_t* sum) {
  67. ReduceOp<acc_t> r;
  68. #pragma unroll
  69. for (int offset = WARP_SIZE / 2; offset > 0; offset /= 2) {
  70. #pragma unroll
  71. for (int i = 0; i < WARP_BATCH; ++i) {
  72. acc_t b = WARP_SHFL_XOR_NATIVE(sum[i], offset, WARP_SIZE);
  73. sum[i] = r(sum[i], b);
  74. }
  75. }
  76. }
  77. /*
  78. * Extended softmax (from native aten pytorch) with following additional features
  79. * 1) input scaling
  80. * 2) Explicit masking
  81. */
  82. template <typename input_t, typename output_t, typename acc_t, int log2_elements>
  83. __global__ void scaled_masked_softmax_warp_forward(
  84. output_t *dst,
  85. const input_t *src,
  86. const uint8_t *mask,
  87. const acc_t scale,
  88. int micro_batch_size,
  89. int element_count,
  90. int pad_batches)
  91. {
  92. // WARP_SIZE and WARP_BATCH must match the return values batches_per_warp and
  93. // warp_size of method warp_softmax_forward_kernel.
  94. constexpr int next_power_of_two = 1 << log2_elements;
  95. constexpr int WARP_SIZE = (next_power_of_two < C10_WARP_SIZE) ? next_power_of_two : C10_WARP_SIZE;
  96. constexpr int WARP_ITERATIONS = next_power_of_two / WARP_SIZE;
  97. constexpr int WARP_BATCH = (next_power_of_two <= 128) ? 2 : 1;
  98. constexpr int ELEMENTS_PER_LDG_STG = 4;
  99. // blockDim/threadIdx = (WARP_SIZE, WARPS_PER_BLOCK, )
  100. // gridDim/blockIdx = (seq_len, attn_heads, batches)
  101. int first_batch = (blockDim.y * (blockIdx.x + gridDim.x * (blockIdx.y + gridDim.y * blockIdx.z))+ threadIdx.y) * WARP_BATCH;
  102. int pad_first_batch = 0;
  103. if (pad_batches != 1) { // bert style
  104. pad_first_batch = (blockDim.y * (blockIdx.x + gridDim.x * blockIdx.z) + threadIdx.y) * WARP_BATCH;
  105. } else { // gpt2 style
  106. pad_first_batch = (blockDim.y * blockIdx.x + threadIdx.y) * WARP_BATCH;
  107. }
  108. // micro_batch_size might not be a multiple of WARP_BATCH. Check how
  109. // many batches have to computed within this WARP.
  110. int local_batches = micro_batch_size - first_batch;
  111. if (local_batches > WARP_BATCH)
  112. local_batches = WARP_BATCH;
  113. // there might be multiple batches per warp. compute the index within the batch
  114. int local_idx = threadIdx.x;
  115. src += first_batch * element_count + ELEMENTS_PER_LDG_STG * local_idx;
  116. dst += first_batch * element_count + ELEMENTS_PER_LDG_STG * local_idx;
  117. mask += pad_first_batch * element_count + ELEMENTS_PER_LDG_STG * local_idx;
  118. // load data from global memory
  119. acc_t elements[WARP_BATCH][WARP_ITERATIONS];
  120. input_t temp_data[ELEMENTS_PER_LDG_STG];
  121. uint8_t temp_mask[ELEMENTS_PER_LDG_STG];
  122. #pragma unroll
  123. for (int i = 0; i < WARP_BATCH; ++i) {
  124. int batch_element_count = (i >= local_batches) ? 0 : element_count;
  125. #pragma unroll
  126. for (int it = 0; it < WARP_ITERATIONS; it+=ELEMENTS_PER_LDG_STG) {
  127. int element_index = ELEMENTS_PER_LDG_STG * local_idx + it * WARP_SIZE;
  128. if (element_index < batch_element_count) {
  129. int itr_idx = i*element_count+it*WARP_SIZE;
  130. copy_vector<input_t, ELEMENTS_PER_LDG_STG>(temp_data, src + itr_idx);
  131. copy_vector<uint8_t, ELEMENTS_PER_LDG_STG>(temp_mask, mask + itr_idx);
  132. #pragma unroll
  133. for (int element = 0; element < ELEMENTS_PER_LDG_STG; ++element) {
  134. if (temp_mask[element] != 1) {
  135. elements[i][it + element] = (acc_t)temp_data[element] * scale;
  136. } else {
  137. elements[i][it + element] = -10000.0;
  138. }
  139. }
  140. } else {
  141. #pragma unroll
  142. for (int element = 0; element < ELEMENTS_PER_LDG_STG; ++element) {
  143. elements[i][it + element] = -std::numeric_limits<acc_t>::infinity();
  144. }
  145. }
  146. }
  147. }
  148. // compute max_value
  149. acc_t max_value[WARP_BATCH];
  150. #pragma unroll
  151. for (int i = 0; i < WARP_BATCH; ++i) {
  152. max_value[i] = elements[i][0];
  153. #pragma unroll
  154. for (int it = 1; it < WARP_ITERATIONS; ++it) {
  155. max_value[i] = (max_value[i] > elements[i][it]) ? max_value[i] : elements[i][it];
  156. }
  157. }
  158. warp_reduce<acc_t, WARP_BATCH, WARP_SIZE, Max>(max_value);
  159. acc_t sum[WARP_BATCH] { 0.0f };
  160. #pragma unroll
  161. for (int i = 0; i < WARP_BATCH; ++i) {
  162. #pragma unroll
  163. for (int it = 0; it < WARP_ITERATIONS; ++it) {
  164. elements[i][it] = std::exp((elements[i][it] - max_value[i]));
  165. sum[i] += elements[i][it];
  166. }
  167. }
  168. warp_reduce<acc_t, WARP_BATCH, WARP_SIZE, Add>(sum);
  169. // store result
  170. output_t out[ELEMENTS_PER_LDG_STG];
  171. #pragma unroll
  172. for (int i = 0; i < WARP_BATCH; ++i) {
  173. if (i >= local_batches)
  174. break;
  175. #pragma unroll
  176. for (int it = 0; it < WARP_ITERATIONS; it+=ELEMENTS_PER_LDG_STG) {
  177. int element_index = ELEMENTS_PER_LDG_STG * local_idx + it * WARP_SIZE;
  178. if (element_index < element_count) {
  179. #pragma unroll
  180. for (int element = 0; element < ELEMENTS_PER_LDG_STG; ++element) {
  181. out[element] = elements[i][it + element] / sum[i];
  182. }
  183. copy_vector<output_t, ELEMENTS_PER_LDG_STG>(dst + i * element_count + it * WARP_SIZE, out);
  184. } else {
  185. break;
  186. }
  187. }
  188. }
  189. }
  190. template <typename input_t, typename output_t, typename acc_t, int log2_elements>
  191. __global__ void scaled_masked_softmax_warp_backward(
  192. output_t *gradInput,
  193. input_t *grad,
  194. const input_t *output,
  195. acc_t scale,
  196. int micro_batch_size,
  197. int element_count)
  198. {
  199. // WARP_SIZE and WARP_BATCH must match the return values batches_per_warp and
  200. // warp_size of method warp_softmax_backward_kernel.
  201. constexpr int next_power_of_two = 1 << log2_elements;
  202. constexpr int WARP_SIZE = (next_power_of_two < C10_WARP_SIZE) ? next_power_of_two : C10_WARP_SIZE;
  203. constexpr int WARP_ITERATIONS = next_power_of_two / WARP_SIZE;
  204. constexpr int WARP_BATCH = (next_power_of_two <= 128) ? 2 : 1;
  205. constexpr int ELEMENTS_PER_LDG_STG = 4;
  206. // blockDim/threadIdx = (WARP_SIZE, WARPS_PER_BLOCK, )
  207. // gridDim/blockIdx = (seq_len, attn_heads, batches)
  208. int first_batch = (blockDim.y * blockIdx.x + threadIdx.y) * WARP_BATCH;
  209. // micro_batch_size might not be a multiple of WARP_BATCH. Check how
  210. // many batches have to computed within this WARP.
  211. int local_batches = micro_batch_size - first_batch;
  212. if (local_batches > WARP_BATCH)
  213. local_batches = WARP_BATCH;
  214. // there might be multiple batches per warp. compute the index within the batch
  215. int local_idx = threadIdx.x;
  216. // the first element to process by the current thread
  217. int thread_offset = first_batch * element_count + ELEMENTS_PER_LDG_STG * local_idx;
  218. grad += thread_offset;
  219. output += thread_offset;
  220. gradInput += thread_offset;
  221. // load data from global memory
  222. acc_t grad_reg[WARP_BATCH][WARP_ITERATIONS] { 0.0f };
  223. acc_t output_reg[WARP_BATCH][WARP_ITERATIONS] { 0.0f };
  224. input_t temp_grad[ELEMENTS_PER_LDG_STG];
  225. input_t temp_output[ELEMENTS_PER_LDG_STG];
  226. #pragma unroll
  227. for (int i = 0; i < WARP_BATCH; ++i) {
  228. int batch_element_count = (i >= local_batches) ? 0 : element_count;
  229. #pragma unroll
  230. for (int it = 0; it < WARP_ITERATIONS; it+=ELEMENTS_PER_LDG_STG) {
  231. int element_index = ELEMENTS_PER_LDG_STG * local_idx + it * WARP_SIZE;
  232. if (element_index < batch_element_count) {
  233. copy_vector<input_t, ELEMENTS_PER_LDG_STG>(temp_grad, grad + i * element_count + it * WARP_SIZE);
  234. copy_vector<input_t, ELEMENTS_PER_LDG_STG>(temp_output, output + i * element_count + it * WARP_SIZE);
  235. #pragma unroll
  236. for (int element = 0; element < ELEMENTS_PER_LDG_STG; ++element) {
  237. output_reg[i][it + element] = (acc_t)temp_output[element];
  238. }
  239. #pragma unroll
  240. for (int element = 0; element < ELEMENTS_PER_LDG_STG; ++element) {
  241. grad_reg[i][it + element] = (acc_t)temp_grad[element] * output_reg[i][it + element];
  242. }
  243. }
  244. }
  245. }
  246. acc_t sum[WARP_BATCH];
  247. #pragma unroll
  248. for (int i = 0; i < WARP_BATCH; ++i) {
  249. sum[i] = grad_reg[i][0];
  250. #pragma unroll
  251. for (int it = 1; it < WARP_ITERATIONS; ++it) {
  252. sum[i] += grad_reg[i][it];
  253. }
  254. }
  255. warp_reduce<acc_t, WARP_BATCH, WARP_SIZE, Add>(sum);
  256. // store result
  257. #pragma unroll
  258. for (int i = 0; i < WARP_BATCH; ++i) {
  259. if (i >= local_batches)
  260. break;
  261. #pragma unroll
  262. for (int it = 0; it < WARP_ITERATIONS; it+=ELEMENTS_PER_LDG_STG) {
  263. int element_index = ELEMENTS_PER_LDG_STG * local_idx + it * WARP_SIZE;
  264. if (element_index < element_count) {
  265. // compute gradients
  266. output_t out[ELEMENTS_PER_LDG_STG];
  267. #pragma unroll
  268. for (int element = 0; element < ELEMENTS_PER_LDG_STG; ++element) {
  269. out[element] = (output_t)(scale * (grad_reg[i][it + element] - output_reg[i][it + element] * sum[i]));
  270. }
  271. copy_vector<output_t, ELEMENTS_PER_LDG_STG>(gradInput + i * element_count + it * WARP_SIZE, out);
  272. }
  273. }
  274. }
  275. }
  276. } // end of anonymous namespace
  277. template<typename input_t, typename output_t, typename acc_t>
  278. void dispatch_scaled_masked_softmax_forward(
  279. output_t *dst,
  280. const input_t *src,
  281. const uint8_t *mask,
  282. const input_t scale,
  283. int query_seq_len,
  284. int key_seq_len,
  285. int batches,
  286. int attn_heads,
  287. int pad_batches)
  288. {
  289. TORCH_INTERNAL_ASSERT(key_seq_len >= 0 && key_seq_len <= 2048 );
  290. if (key_seq_len == 0) {
  291. return;
  292. } else {
  293. int log2_elements = log2_ceil(key_seq_len);
  294. const int next_power_of_two = 1 << log2_elements;
  295. int batch_count = batches * attn_heads * query_seq_len;
  296. // This value must match the WARP_SIZE constexpr value computed inside softmax_warp_forward.
  297. int warp_size = (next_power_of_two < C10_WARP_SIZE) ? next_power_of_two : C10_WARP_SIZE;
  298. // This value must match the WARP_BATCH constexpr value computed inside softmax_warp_forward.
  299. int batches_per_warp = (next_power_of_two <= 128) ? 2 : 1;
  300. // use 128 threads per block to maximimize gpu utilization
  301. constexpr int threads_per_block = 128;
  302. int warps_per_block = (threads_per_block / warp_size);
  303. int batches_per_block = warps_per_block * batches_per_warp;
  304. TORCH_INTERNAL_ASSERT(query_seq_len%batches_per_block == 0);
  305. dim3 blocks(query_seq_len/batches_per_block, attn_heads, batches);
  306. dim3 threads(warp_size, warps_per_block, 1);
  307. // Launch code would be more elegant if C++ supported FOR CONSTEXPR
  308. switch (log2_elements) {
  309. case 0: // 1
  310. scaled_masked_softmax_warp_forward<input_t, output_t, acc_t, 0>
  311. <<<blocks, threads, 0, at::cuda::getCurrentCUDAStream()>>>(dst, src, mask, scale, batch_count, key_seq_len, pad_batches);
  312. break;
  313. case 1: // 2
  314. scaled_masked_softmax_warp_forward<input_t, output_t, acc_t, 1>
  315. <<<blocks, threads, 0, at::cuda::getCurrentCUDAStream()>>>(dst, src, mask, scale, batch_count, key_seq_len, pad_batches);
  316. break;
  317. case 2: // 4
  318. scaled_masked_softmax_warp_forward<input_t, output_t, acc_t, 2>
  319. <<<blocks, threads, 0, at::cuda::getCurrentCUDAStream()>>>(dst, src, mask, scale, batch_count, key_seq_len, pad_batches);
  320. break;
  321. case 3: // 8
  322. scaled_masked_softmax_warp_forward<input_t, output_t, acc_t, 3>
  323. <<<blocks, threads, 0, at::cuda::getCurrentCUDAStream()>>>(dst, src, mask, scale, batch_count, key_seq_len, pad_batches);
  324. break;
  325. case 4: // 16
  326. scaled_masked_softmax_warp_forward<input_t, output_t, acc_t, 4>
  327. <<<blocks, threads, 0, at::cuda::getCurrentCUDAStream()>>>(dst, src, mask, scale, batch_count, key_seq_len, pad_batches);
  328. break;
  329. case 5: // 32
  330. scaled_masked_softmax_warp_forward<input_t, output_t, acc_t, 5>
  331. <<<blocks, threads, 0, at::cuda::getCurrentCUDAStream()>>>(dst, src, mask, scale, batch_count, key_seq_len, pad_batches);
  332. break;
  333. case 6: // 64
  334. scaled_masked_softmax_warp_forward<input_t, output_t, acc_t, 6>
  335. <<<blocks, threads, 0, at::cuda::getCurrentCUDAStream()>>>(dst, src, mask, scale, batch_count, key_seq_len, pad_batches);
  336. break;
  337. case 7: // 128
  338. scaled_masked_softmax_warp_forward<input_t, output_t, acc_t, 7>
  339. <<<blocks, threads, 0, at::cuda::getCurrentCUDAStream()>>>(dst, src, mask, scale, batch_count, key_seq_len, pad_batches);
  340. break;
  341. case 8: // 256
  342. scaled_masked_softmax_warp_forward<input_t, output_t, acc_t, 8>
  343. <<<blocks, threads, 0, at::cuda::getCurrentCUDAStream()>>>(dst, src, mask, scale, batch_count, key_seq_len, pad_batches);
  344. break;
  345. case 9: // 512
  346. scaled_masked_softmax_warp_forward<input_t, output_t, acc_t, 9>
  347. <<<blocks, threads, 0, at::cuda::getCurrentCUDAStream()>>>(dst, src, mask, scale, batch_count, key_seq_len, pad_batches);
  348. break;
  349. case 10: // 1024
  350. scaled_masked_softmax_warp_forward<input_t, output_t, acc_t, 10>
  351. <<<blocks, threads, 0, at::cuda::getCurrentCUDAStream()>>>(dst, src, mask, scale, batch_count, key_seq_len, pad_batches);
  352. break;
  353. case 11: // 2048
  354. scaled_masked_softmax_warp_forward<input_t, output_t, acc_t, 11>
  355. <<<blocks, threads, 0, at::cuda::getCurrentCUDAStream()>>>(dst, src, mask, scale, batch_count, key_seq_len, pad_batches);
  356. break;
  357. default:
  358. break;
  359. }
  360. }
  361. }
  362. template<typename input_t, typename output_t, typename acc_t>
  363. void dispatch_scaled_masked_softmax_backward(
  364. output_t *grad_input,
  365. input_t *grad,
  366. const input_t *output,
  367. const acc_t scale,
  368. int query_seq_len,
  369. int key_seq_len,
  370. int batches,
  371. int attn_heads)
  372. {
  373. TORCH_INTERNAL_ASSERT( key_seq_len >= 0 && key_seq_len <= 2048 );
  374. if (key_seq_len == 0) {
  375. return;
  376. } else {
  377. int log2_elements = log2_ceil(key_seq_len);
  378. const int next_power_of_two = 1 << log2_elements;
  379. int batch_count = batches * attn_heads * query_seq_len;
  380. // This value must match the WARP_SIZE constexpr value computed inside softmax_warp_backward.
  381. int warp_size = (next_power_of_two < C10_WARP_SIZE) ? next_power_of_two : C10_WARP_SIZE;
  382. // This value must match the WARP_BATCH constexpr value computed inside softmax_warp_backward.
  383. int batches_per_warp = (next_power_of_two <= 128) ? 2 : 1;
  384. // use 128 threads per block to maximimize gpu utilization
  385. constexpr int threads_per_block = 128;
  386. int warps_per_block = (threads_per_block / warp_size);
  387. int batches_per_block = warps_per_block * batches_per_warp;
  388. int blocks = batch_count/batches_per_block;
  389. dim3 threads(warp_size, warps_per_block, 1);
  390. // Launch code would be more elegant if C++ supported FOR CONSTEXPR
  391. switch (log2_elements) {
  392. case 0: // 1
  393. scaled_masked_softmax_warp_backward<input_t, output_t, acc_t, 0>
  394. <<<blocks, threads, 0, at::cuda::getCurrentCUDAStream()>>>(grad_input, grad, output, scale, batch_count, key_seq_len);
  395. break;
  396. case 1: // 2
  397. scaled_masked_softmax_warp_backward<input_t, output_t, acc_t, 1>
  398. <<<blocks, threads, 0, at::cuda::getCurrentCUDAStream()>>>(grad_input, grad, output, scale, batch_count, key_seq_len);
  399. break;
  400. case 2: // 4
  401. scaled_masked_softmax_warp_backward<input_t, output_t, acc_t, 2>
  402. <<<blocks, threads, 0, at::cuda::getCurrentCUDAStream()>>>(grad_input, grad, output, scale, batch_count, key_seq_len);
  403. break;
  404. case 3: // 8
  405. scaled_masked_softmax_warp_backward<input_t, output_t, acc_t, 3>
  406. <<<blocks, threads, 0, at::cuda::getCurrentCUDAStream()>>>(grad_input, grad, output, scale, batch_count, key_seq_len);
  407. break;
  408. case 4: // 16
  409. scaled_masked_softmax_warp_backward<input_t, output_t, acc_t, 4>
  410. <<<blocks, threads, 0, at::cuda::getCurrentCUDAStream()>>>(grad_input, grad, output, scale, batch_count, key_seq_len);
  411. break;
  412. case 5: // 32
  413. scaled_masked_softmax_warp_backward<input_t, output_t, acc_t, 5>
  414. <<<blocks, threads, 0, at::cuda::getCurrentCUDAStream()>>>(grad_input, grad, output, scale, batch_count, key_seq_len);
  415. break;
  416. case 6: // 64
  417. scaled_masked_softmax_warp_backward<input_t, output_t, acc_t, 6>
  418. <<<blocks, threads, 0, at::cuda::getCurrentCUDAStream()>>>(grad_input, grad, output, scale, batch_count, key_seq_len);
  419. break;
  420. case 7: // 128
  421. scaled_masked_softmax_warp_backward<input_t, output_t, acc_t, 7>
  422. <<<blocks, threads, 0, at::cuda::getCurrentCUDAStream()>>>(grad_input, grad, output, scale, batch_count, key_seq_len);
  423. break;
  424. case 8: // 256
  425. scaled_masked_softmax_warp_backward<input_t, output_t, acc_t, 8>
  426. <<<blocks, threads, 0, at::cuda::getCurrentCUDAStream()>>>(grad_input, grad, output, scale, batch_count, key_seq_len);
  427. break;
  428. case 9: // 512
  429. scaled_masked_softmax_warp_backward<input_t, output_t, acc_t, 9>
  430. <<<blocks, threads, 0, at::cuda::getCurrentCUDAStream()>>>(grad_input, grad, output, scale, batch_count, key_seq_len);
  431. break;
  432. case 10: // 1024
  433. scaled_masked_softmax_warp_backward<input_t, output_t, acc_t, 10>
  434. <<<blocks, threads, 0, at::cuda::getCurrentCUDAStream()>>>(grad_input, grad, output, scale, batch_count, key_seq_len);
  435. break;
  436. case 11: // 2048
  437. scaled_masked_softmax_warp_backward<input_t, output_t, acc_t, 11>
  438. <<<blocks, threads, 0, at::cuda::getCurrentCUDAStream()>>>(grad_input, grad, output, scale, batch_count, key_seq_len);
  439. break;
  440. default:
  441. break;
  442. }
  443. }
  444. }