scaled_upper_triang_masked_softmax.h 24 KB

123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121122123124125126127128129130131132133134135136137138139140141142143144145146147148149150151152153154155156157158159160161162163164165166167168169170171172173174175176177178179180181182183184185186187188189190191192193194195196197198199200201202203204205206207208209210211212213214215216217218219220221222223224225226227228229230231232233234235236237238239240241242243244245246247248249250251252253254255256257258259260261262263264265266267268269270271272273274275276277278279280281282283284285286287288289290291292293294295296297298299300301302303304305306307308309310311312313314315316317318319320321322323324325326327328329330331332333334335336337338339340341342343344345346347348349350351352353354355356357358359360361362363364365366367368369370371372373374375376377378379380381382383384385386387388389390391392393394395396397398399400401402403404405406407408409410411412413414415416417418419420421422423424425426427428429430431432433434435436437438439440441442443444445446447448449450451452453454455456457458459460461462463464465466467468469470471472473474475476477478479480481482483484485486487488489490491492493494495496497498499500501502503504505506507508509510511512513514515516517518519520521522523524525526527528529
  1. /* coding=utf-8
  2. * Copyright (c) 2021, 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 <c10/macros/Macros.h>
  23. namespace {
  24. template <typename Datatype, int ELEMENTS_PER_LDG>
  25. __device__ __inline__ void copy_vector(Datatype *dst, const Datatype *src);
  26. template <>
  27. __device__ __inline__ void copy_vector<c10::BFloat16, 1>(c10::BFloat16 *dst, const c10::BFloat16 *src) { *dst = *src; }
  28. template <>
  29. __device__ __inline__ void copy_vector<c10::BFloat16, 4>(c10::BFloat16 *dst, const c10::BFloat16 *src) { *((float2*) dst) = *((float2*) src); }
  30. template <>
  31. __device__ __inline__ void copy_vector<c10::Half, 1>(c10::Half *dst, const c10::Half *src) { *dst = *src; }
  32. template <>
  33. __device__ __inline__ void copy_vector<c10::Half, 4>(c10::Half *dst, const c10::Half *src) { *((float2*) dst) = *((float2*) src); }
  34. template <>
  35. __device__ __inline__ void copy_vector<uint8_t, 1>(uint8_t *dst, const uint8_t *src) { *dst = *src; }
  36. template <>
  37. __device__ __inline__ void copy_vector<uint8_t, 4>(uint8_t *dst, const uint8_t *src) {*((half2*) dst) = *((half2*) src); }
  38. template <typename Datatype, int ELEMENTS_PER_LDG>
  39. __device__ __inline__ void copy_zero_vector(Datatype *dst);
  40. template <>
  41. __device__ __inline__ void copy_zero_vector<c10::BFloat16, 1>(c10::BFloat16 *dst) { *dst = 0.0; }
  42. template <>
  43. __device__ __inline__ void copy_zero_vector<c10::BFloat16, 4>(c10::BFloat16 *dst) { *((float2*) dst) = make_float2(0.0f, 0.0f); }
  44. template <>
  45. __device__ __inline__ void copy_zero_vector<c10::Half, 1>(c10::Half *dst) { *dst = 0.0; }
  46. template <>
  47. __device__ __inline__ void copy_zero_vector<c10::Half, 4>(c10::Half *dst) { *((float2*) dst) = make_float2(0.0f, 0.0f); }
  48. int log2_ceil(int value) {
  49. int log2_value = 0;
  50. while ((1 << log2_value) < value) ++log2_value;
  51. return log2_value;
  52. }
  53. template<typename T>
  54. struct Add {
  55. __device__ __forceinline__ T operator()(T a, T b) const {
  56. return a + b;
  57. }
  58. };
  59. template<typename T>
  60. struct Max {
  61. __device__ __forceinline__ T operator()(T a, T b) const {
  62. return a < b ? b : a;
  63. }
  64. };
  65. template <typename T>
  66. __device__ __forceinline__ T WARP_SHFL_XOR_NATIVE(T value, int laneMask, int width = warpSize, unsigned int mask = 0xffffffff)
  67. {
  68. #if CUDA_VERSION >= 9000
  69. return __shfl_xor_sync(mask, value, laneMask, width);
  70. #else
  71. return __shfl_xor(value, laneMask, width);
  72. #endif
  73. }
  74. template <typename acc_t, int WARP_BATCH, int WARP_SIZE, template<typename> class ReduceOp>
  75. __device__ __forceinline__ void warp_reduce(acc_t* sum) {
  76. ReduceOp<acc_t> r;
  77. #pragma unroll
  78. for (int offset = WARP_SIZE / 2; offset > 0; offset /= 2) {
  79. #pragma unroll
  80. for (int i = 0; i < WARP_BATCH; ++i) {
  81. acc_t b = WARP_SHFL_XOR_NATIVE(sum[i], offset, WARP_SIZE);
  82. sum[i] = r(sum[i], b);
  83. }
  84. }
  85. }
  86. /*
  87. * Extended softmax (from native aten pytorch) with following additional features
  88. * 1) input scaling
  89. * 2) Implicit time (diagonal masking)
  90. */
  91. template <typename input_t, typename output_t, typename acc_t, int log2_elements>
  92. __global__ void scaled_upper_triang_masked_softmax_warp_forward(
  93. output_t *dst,
  94. const input_t *src,
  95. const acc_t scale,
  96. int micro_batch_size,
  97. int stride,
  98. int element_count)
  99. {
  100. // WARP_SIZE and WARP_BATCH must match the return values batches_per_warp and
  101. // warp_size of method warp_softmax_forward_kernel.
  102. constexpr int next_power_of_two = 1 << log2_elements;
  103. constexpr int WARP_SIZE = (next_power_of_two < C10_WARP_SIZE) ? next_power_of_two : C10_WARP_SIZE;
  104. constexpr int WARP_ITERATIONS = next_power_of_two / WARP_SIZE;
  105. constexpr int WARP_BATCH = (next_power_of_two <= 128) ? 2 : 1;
  106. constexpr int ELEMENTS_PER_LDG_STG = (WARP_ITERATIONS < 4) ? 1 : 4;
  107. int first_batch = (blockDim.y * blockIdx.y + threadIdx.y) * gridDim.x * WARP_BATCH + blockIdx.x;
  108. int local_seq = blockIdx.x + 1;
  109. int warp_iteration_limit = (local_seq + ELEMENTS_PER_LDG_STG * WARP_SIZE - 1)/ WARP_SIZE;
  110. // micro_batch_size might not be a multiple of WARP_BATCH. Check how
  111. // many batches have to computed within this WARP.
  112. int local_batches = micro_batch_size - first_batch;
  113. if (local_batches > WARP_BATCH)
  114. local_batches = WARP_BATCH;
  115. // there might be multiple batches per warp. compute the index within the batch
  116. int local_idx = threadIdx.x;
  117. src += first_batch * stride + ELEMENTS_PER_LDG_STG * local_idx;
  118. dst += first_batch * stride + ELEMENTS_PER_LDG_STG * local_idx;
  119. // load data from global memory
  120. acc_t elements[WARP_BATCH][WARP_ITERATIONS];
  121. input_t temp_data[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 : local_seq;
  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. copy_vector<input_t, ELEMENTS_PER_LDG_STG>(temp_data, src + i*element_count*stride + it*WARP_SIZE);
  130. #pragma unroll
  131. for (int element = 0; element < ELEMENTS_PER_LDG_STG; ++element) {
  132. if ((element_index + element) < batch_element_count) {
  133. elements[i][it+element] = (acc_t)temp_data[element] * scale;
  134. } else {
  135. elements[i][it + element] = -std::numeric_limits<acc_t>::infinity();
  136. }
  137. }
  138. } else {
  139. #pragma unroll
  140. for (int element = 0; element < ELEMENTS_PER_LDG_STG; ++element) {
  141. elements[i][it + element] = -std::numeric_limits<acc_t>::infinity();
  142. }
  143. }
  144. }
  145. }
  146. // compute max_value
  147. acc_t max_value[WARP_BATCH];
  148. #pragma unroll
  149. for (int i = 0; i < WARP_BATCH; ++i) {
  150. max_value[i] = elements[i][0];
  151. #pragma unroll
  152. for (int it = 1; it < WARP_ITERATIONS; ++it) {
  153. max_value[i] = (max_value[i] > elements[i][it]) ? max_value[i] : elements[i][it];
  154. }
  155. }
  156. warp_reduce<acc_t, WARP_BATCH, WARP_SIZE, Max>(max_value);
  157. acc_t sum[WARP_BATCH] { 0.0f };
  158. #pragma unroll
  159. for (int i = 0; i < WARP_BATCH; ++i) {
  160. #pragma unroll
  161. for (int it = 0; it < WARP_ITERATIONS; ++it) {
  162. if (it < warp_iteration_limit) {
  163. elements[i][it] = std::exp((elements[i][it] - max_value[i]));
  164. sum[i] += elements[i][it];
  165. }
  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 < local_seq) {
  179. #pragma unroll
  180. for (int element = 0; element < ELEMENTS_PER_LDG_STG; ++element) {
  181. if (element_index + element < local_seq) {
  182. out[element] = elements[i][it + element] / sum[i];
  183. } else {
  184. out[element] = 0;
  185. }
  186. }
  187. copy_vector<output_t, ELEMENTS_PER_LDG_STG>(dst + i * element_count * stride + it * WARP_SIZE, out);
  188. } else if (element_index < element_count) {
  189. copy_zero_vector<output_t, ELEMENTS_PER_LDG_STG>(dst + i * element_count * stride + it * WARP_SIZE);
  190. } else {
  191. break;
  192. }
  193. }
  194. }
  195. }
  196. template <typename input_t, typename output_t, typename acc_t, int log2_elements>
  197. __global__ void scaled_upper_triang_masked_softmax_warp_backward(
  198. output_t *gradInput,
  199. input_t *grad,
  200. const input_t *output,
  201. acc_t scale,
  202. int micro_batch_size,
  203. int stride,
  204. int element_count)
  205. {
  206. // WARP_SIZE and WARP_BATCH must match the return values batches_per_warp and
  207. // warp_size of method warp_softmax_backward_kernel.
  208. constexpr int next_power_of_two = 1 << log2_elements;
  209. constexpr int WARP_SIZE = (next_power_of_two < C10_WARP_SIZE) ? next_power_of_two : C10_WARP_SIZE;
  210. constexpr int WARP_ITERATIONS = next_power_of_two / WARP_SIZE;
  211. constexpr int WARP_BATCH = (next_power_of_two <= 128) ? 2 : 1;
  212. constexpr int ELEMENTS_PER_LDG_STG = (WARP_ITERATIONS < 4) ? 1 : 4;
  213. int first_batch = (blockDim.y * blockIdx.y + threadIdx.y) * gridDim.x * WARP_BATCH + blockIdx.x;
  214. int local_seq = blockIdx.x + 1;
  215. // micro_batch_size might not be a multiple of WARP_BATCH. Check how
  216. // many batches have to computed within this WARP.
  217. int local_batches = micro_batch_size - first_batch;
  218. if (local_batches > WARP_BATCH)
  219. local_batches = WARP_BATCH;
  220. // there might be multiple batches per warp. compute the index within the batch
  221. int local_idx = threadIdx.x;
  222. // the first element to process by the current thread
  223. int thread_offset = first_batch * stride + ELEMENTS_PER_LDG_STG * local_idx;
  224. grad += thread_offset;
  225. output += thread_offset;
  226. gradInput += thread_offset;
  227. // load data from global memory
  228. acc_t grad_reg[WARP_BATCH][WARP_ITERATIONS] { 0.0f };
  229. acc_t output_reg[WARP_BATCH][WARP_ITERATIONS] { 0.0f };
  230. input_t temp_grad[ELEMENTS_PER_LDG_STG];
  231. input_t temp_output[ELEMENTS_PER_LDG_STG];
  232. #pragma unroll
  233. for (int i = 0; i < WARP_BATCH; ++i) {
  234. int batch_element_count = (i >= local_batches) ? 0 : local_seq;
  235. #pragma unroll
  236. for (int it = 0; it < WARP_ITERATIONS; it+=ELEMENTS_PER_LDG_STG) {
  237. int element_index = ELEMENTS_PER_LDG_STG * local_idx + it * WARP_SIZE;
  238. if (element_index < batch_element_count) {
  239. copy_vector<input_t, ELEMENTS_PER_LDG_STG>(temp_grad, grad + i * element_count * stride + it * WARP_SIZE);
  240. copy_vector<input_t, ELEMENTS_PER_LDG_STG>(temp_output, output + i * element_count * stride + it * WARP_SIZE);
  241. #pragma unroll
  242. for (int element = 0; element < ELEMENTS_PER_LDG_STG; ++element) {
  243. if (element_index + element < batch_element_count) {
  244. output_reg[i][it + element] = (acc_t)temp_output[element];
  245. }
  246. }
  247. #pragma unroll
  248. for (int element = 0; element < ELEMENTS_PER_LDG_STG; ++element) {
  249. if (element_index + element < batch_element_count) {
  250. grad_reg[i][it + element] = (acc_t)temp_grad[element] * output_reg[i][it + element];
  251. }
  252. }
  253. }
  254. }
  255. }
  256. acc_t sum[WARP_BATCH];
  257. #pragma unroll
  258. for (int i = 0; i < WARP_BATCH; ++i) {
  259. sum[i] = grad_reg[i][0];
  260. #pragma unroll
  261. for (int it = 1; it < WARP_ITERATIONS; ++it) {
  262. sum[i] += grad_reg[i][it];
  263. }
  264. }
  265. warp_reduce<acc_t, WARP_BATCH, WARP_SIZE, Add>(sum);
  266. // store result
  267. #pragma unroll
  268. for (int i = 0; i < WARP_BATCH; ++i) {
  269. if (i >= local_batches)
  270. break;
  271. #pragma unroll
  272. for (int it = 0; it < WARP_ITERATIONS; it+=ELEMENTS_PER_LDG_STG) {
  273. int element_index = ELEMENTS_PER_LDG_STG * local_idx + it * WARP_SIZE;
  274. if (element_index < element_count) {
  275. // compute gradients
  276. output_t out[ELEMENTS_PER_LDG_STG];
  277. #pragma unroll
  278. for (int element = 0; element < ELEMENTS_PER_LDG_STG; ++element) {
  279. out[element] = (output_t)(scale * (grad_reg[i][it + element] - output_reg[i][it + element] * sum[i]));
  280. }
  281. copy_vector<output_t, ELEMENTS_PER_LDG_STG>(gradInput + i * element_count * stride + it * WARP_SIZE, out);
  282. }
  283. }
  284. }
  285. }
  286. } // end of anonymous namespace
  287. template<typename input_t, typename output_t, typename acc_t>
  288. void dispatch_scaled_upper_triang_masked_softmax_forward(
  289. output_t *dst,
  290. const input_t *src,
  291. const input_t scale,
  292. int softmax_elements,
  293. int softmax_elements_stride,
  294. int attn_batches)
  295. {
  296. TORCH_INTERNAL_ASSERT(softmax_elements >= 0 && softmax_elements <= 8192 );
  297. if (softmax_elements == 0) {
  298. return;
  299. } else {
  300. int log2_elements = log2_ceil(softmax_elements);
  301. const int next_power_of_two = 1 << log2_elements;
  302. int seq_len = softmax_elements;
  303. int batch_count = attn_batches * seq_len;
  304. // This value must match the WARP_SIZE constexpr value computed inside softmax_warp_forward.
  305. int warp_size = (next_power_of_two < C10_WARP_SIZE) ? next_power_of_two : C10_WARP_SIZE;
  306. // This value must match the WARP_BATCH constexpr value computed inside softmax_warp_forward.
  307. int batches_per_warp = (next_power_of_two <= 128) ? 2 : 1;
  308. // use 128 threads per block to maximimize gpu utilization
  309. constexpr int threads_per_block = 128;
  310. int warps_per_block = (threads_per_block / warp_size);
  311. int batches_per_block = warps_per_block * batches_per_warp;
  312. TORCH_INTERNAL_ASSERT(attn_batches % batches_per_block == 0);
  313. int blocks_per_seq = attn_batches / batches_per_block;
  314. dim3 blocks(seq_len, blocks_per_seq, 1);
  315. dim3 threads(warp_size, warps_per_block, 1);
  316. // Launch code would be more elegant if C++ supported FOR CONSTEXPR
  317. switch (log2_elements) {
  318. case 0: // 1
  319. scaled_upper_triang_masked_softmax_warp_forward<input_t, output_t, acc_t, 0>
  320. <<<blocks, threads, 0, at::cuda::getCurrentCUDAStream()>>>(dst, src, scale, batch_count, softmax_elements_stride, softmax_elements);
  321. break;
  322. case 1: // 2
  323. scaled_upper_triang_masked_softmax_warp_forward<input_t, output_t, acc_t, 1>
  324. <<<blocks, threads, 0, at::cuda::getCurrentCUDAStream()>>>(dst, src, scale, batch_count, softmax_elements_stride, softmax_elements);
  325. break;
  326. case 2: // 4
  327. scaled_upper_triang_masked_softmax_warp_forward<input_t, output_t, acc_t, 2>
  328. <<<blocks, threads, 0, at::cuda::getCurrentCUDAStream()>>>(dst, src, scale, batch_count, softmax_elements_stride, softmax_elements);
  329. break;
  330. case 3: // 8
  331. scaled_upper_triang_masked_softmax_warp_forward<input_t, output_t, acc_t, 3>
  332. <<<blocks, threads, 0, at::cuda::getCurrentCUDAStream()>>>(dst, src, scale, batch_count, softmax_elements_stride, softmax_elements);
  333. break;
  334. case 4: // 16
  335. scaled_upper_triang_masked_softmax_warp_forward<input_t, output_t, acc_t, 4>
  336. <<<blocks, threads, 0, at::cuda::getCurrentCUDAStream()>>>(dst, src, scale, batch_count, softmax_elements_stride, softmax_elements);
  337. break;
  338. case 5: // 32
  339. scaled_upper_triang_masked_softmax_warp_forward<input_t, output_t, acc_t, 5>
  340. <<<blocks, threads, 0, at::cuda::getCurrentCUDAStream()>>>(dst, src, scale, batch_count, softmax_elements_stride, softmax_elements);
  341. break;
  342. case 6: // 64
  343. scaled_upper_triang_masked_softmax_warp_forward<input_t, output_t, acc_t, 6>
  344. <<<blocks, threads, 0, at::cuda::getCurrentCUDAStream()>>>(dst, src, scale, batch_count, softmax_elements_stride, softmax_elements);
  345. break;
  346. case 7: // 128
  347. scaled_upper_triang_masked_softmax_warp_forward<input_t, output_t, acc_t, 7>
  348. <<<blocks, threads, 0, at::cuda::getCurrentCUDAStream()>>>(dst, src, scale, batch_count, softmax_elements_stride, softmax_elements);
  349. break;
  350. case 8: // 256
  351. scaled_upper_triang_masked_softmax_warp_forward<input_t, output_t, acc_t, 8>
  352. <<<blocks, threads, 0, at::cuda::getCurrentCUDAStream()>>>(dst, src, scale, batch_count, softmax_elements_stride, softmax_elements);
  353. break;
  354. case 9: // 512
  355. scaled_upper_triang_masked_softmax_warp_forward<input_t, output_t, acc_t, 9>
  356. <<<blocks, threads, 0, at::cuda::getCurrentCUDAStream()>>>(dst, src, scale, batch_count, softmax_elements_stride, softmax_elements);
  357. break;
  358. case 10: // 1024
  359. scaled_upper_triang_masked_softmax_warp_forward<input_t, output_t, acc_t, 10>
  360. <<<blocks, threads, 0, at::cuda::getCurrentCUDAStream()>>>(dst, src, scale, batch_count, softmax_elements_stride, softmax_elements);
  361. break;
  362. case 11: // 2048
  363. scaled_upper_triang_masked_softmax_warp_forward<input_t, output_t, acc_t, 11>
  364. <<<blocks, threads, 0, at::cuda::getCurrentCUDAStream()>>>(dst, src, scale, batch_count, softmax_elements_stride, softmax_elements);
  365. break;
  366. case 12: // 4096
  367. scaled_upper_triang_masked_softmax_warp_forward<input_t, output_t, acc_t, 12>
  368. <<<blocks, threads, 0, at::cuda::getCurrentCUDAStream()>>>(dst, src, scale, batch_count, softmax_elements_stride, softmax_elements);
  369. break;
  370. case 13: // 8192
  371. scaled_upper_triang_masked_softmax_warp_forward<input_t, output_t, acc_t, 13>
  372. <<<blocks, threads, 0, at::cuda::getCurrentCUDAStream()>>>(dst, src, scale, batch_count, softmax_elements_stride, softmax_elements);
  373. break;
  374. default:
  375. break;
  376. }
  377. }
  378. }
  379. template<typename input_t, typename output_t, typename acc_t>
  380. void dispatch_scaled_upper_triang_masked_softmax_backward(
  381. output_t *grad_input,
  382. input_t *grad,
  383. const input_t *output,
  384. const acc_t scale,
  385. int softmax_elements,
  386. int softmax_elements_stride,
  387. int attn_batches)
  388. {
  389. TORCH_INTERNAL_ASSERT( softmax_elements >= 0 && softmax_elements <= 8192 );
  390. if (softmax_elements == 0) {
  391. return;
  392. } else {
  393. int log2_elements = log2_ceil(softmax_elements);
  394. const int next_power_of_two = 1 << log2_elements;
  395. int seq_len = softmax_elements;
  396. int batch_count = attn_batches * seq_len;
  397. // This value must match the WARP_SIZE constexpr value computed inside softmax_warp_backward.
  398. int warp_size = (next_power_of_two < C10_WARP_SIZE) ? next_power_of_two : C10_WARP_SIZE;
  399. // This value must match the WARP_BATCH constexpr value computed inside softmax_warp_backward.
  400. int batches_per_warp = (next_power_of_two <= 128) ? 2 : 1;
  401. // use 128 threads per block to maximimize gpu utilization
  402. constexpr int threads_per_block = 128;
  403. int warps_per_block = (threads_per_block / warp_size);
  404. int batches_per_block = warps_per_block * batches_per_warp;
  405. TORCH_INTERNAL_ASSERT(attn_batches % batches_per_block == 0);
  406. int blocks_per_seq = attn_batches / batches_per_block;
  407. dim3 blocks(seq_len, blocks_per_seq, 1);
  408. dim3 threads(warp_size, warps_per_block, 1);
  409. // Launch code would be more elegant if C++ supported FOR CONSTEXPR
  410. switch (log2_elements) {
  411. case 0: // 1
  412. scaled_upper_triang_masked_softmax_warp_backward<input_t, output_t, acc_t, 0>
  413. <<<blocks, threads, 0, at::cuda::getCurrentCUDAStream()>>>(grad_input, grad, output, scale, batch_count, softmax_elements_stride, softmax_elements);
  414. break;
  415. case 1: // 2
  416. scaled_upper_triang_masked_softmax_warp_backward<input_t, output_t, acc_t, 1>
  417. <<<blocks, threads, 0, at::cuda::getCurrentCUDAStream()>>>(grad_input, grad, output, scale, batch_count, softmax_elements_stride, softmax_elements);
  418. break;
  419. case 2: // 4
  420. scaled_upper_triang_masked_softmax_warp_backward<input_t, output_t, acc_t, 2>
  421. <<<blocks, threads, 0, at::cuda::getCurrentCUDAStream()>>>(grad_input, grad, output, scale, batch_count, softmax_elements_stride, softmax_elements);
  422. break;
  423. case 3: // 8
  424. scaled_upper_triang_masked_softmax_warp_backward<input_t, output_t, acc_t, 3>
  425. <<<blocks, threads, 0, at::cuda::getCurrentCUDAStream()>>>(grad_input, grad, output, scale, batch_count, softmax_elements_stride, softmax_elements);
  426. break;
  427. case 4: // 16
  428. scaled_upper_triang_masked_softmax_warp_backward<input_t, output_t, acc_t, 4>
  429. <<<blocks, threads, 0, at::cuda::getCurrentCUDAStream()>>>(grad_input, grad, output, scale, batch_count, softmax_elements_stride, softmax_elements);
  430. break;
  431. case 5: // 32
  432. scaled_upper_triang_masked_softmax_warp_backward<input_t, output_t, acc_t, 5>
  433. <<<blocks, threads, 0, at::cuda::getCurrentCUDAStream()>>>(grad_input, grad, output, scale, batch_count, softmax_elements_stride, softmax_elements);
  434. break;
  435. case 6: // 64
  436. scaled_upper_triang_masked_softmax_warp_backward<input_t, output_t, acc_t, 6>
  437. <<<blocks, threads, 0, at::cuda::getCurrentCUDAStream()>>>(grad_input, grad, output, scale, batch_count, softmax_elements_stride, softmax_elements);
  438. break;
  439. case 7: // 128
  440. scaled_upper_triang_masked_softmax_warp_backward<input_t, output_t, acc_t, 7>
  441. <<<blocks, threads, 0, at::cuda::getCurrentCUDAStream()>>>(grad_input, grad, output, scale, batch_count, softmax_elements_stride, softmax_elements);
  442. break;
  443. case 8: // 256
  444. scaled_upper_triang_masked_softmax_warp_backward<input_t, output_t, acc_t, 8>
  445. <<<blocks, threads, 0, at::cuda::getCurrentCUDAStream()>>>(grad_input, grad, output, scale, batch_count, softmax_elements_stride, softmax_elements);
  446. break;
  447. case 9: // 512
  448. scaled_upper_triang_masked_softmax_warp_backward<input_t, output_t, acc_t, 9>
  449. <<<blocks, threads, 0, at::cuda::getCurrentCUDAStream()>>>(grad_input, grad, output, scale, batch_count, softmax_elements_stride, softmax_elements);
  450. break;
  451. case 10: // 1024
  452. scaled_upper_triang_masked_softmax_warp_backward<input_t, output_t, acc_t, 10>
  453. <<<blocks, threads, 0, at::cuda::getCurrentCUDAStream()>>>(grad_input, grad, output, scale, batch_count, softmax_elements_stride, softmax_elements);
  454. break;
  455. case 11: // 2048
  456. scaled_upper_triang_masked_softmax_warp_backward<input_t, output_t, acc_t, 11>
  457. <<<blocks, threads, 0, at::cuda::getCurrentCUDAStream()>>>(grad_input, grad, output, scale, batch_count, softmax_elements_stride, softmax_elements);
  458. break;
  459. case 12: // 4096
  460. scaled_upper_triang_masked_softmax_warp_backward<input_t, output_t, acc_t, 12>
  461. <<<blocks, threads, 0, at::cuda::getCurrentCUDAStream()>>>(grad_input, grad, output, scale, batch_count, softmax_elements_stride, softmax_elements);
  462. break;
  463. case 13: // 8192
  464. scaled_upper_triang_masked_softmax_warp_backward<input_t, output_t, acc_t, 13>
  465. <<<blocks, threads, 0, at::cuda::getCurrentCUDAStream()>>>(grad_input, grad, output, scale, batch_count, softmax_elements_stride, softmax_elements);
  466. break;
  467. default:
  468. break;
  469. }
  470. }
  471. }