ln_api.cpp 20 KB

123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121122123124125126127128129130131132133134135136137138139140141142143144145146147148149150151152153154155156157158159160161162163164165166167168169170171172173174175176177178179180181182183184185186187188189190191192193194195196197198199200201202203204205206207208209210211212213214215216217218219220221222223224225226227228229230231232233234235236237238239240241242243244245246247248249250251252253254255256257258259260261262263264265266267268269270271272273274275276277278279280281282283284285286287288289290291292293294295296297298299300301302303304305306307308309310311312313314315316317318319320321322323324325326327328329330331332333334335336337338339340341342343344345346347348349350351352353354355356357358359360361362363364365366367368369370371372373374375376377378379380381382383384385386387388389390391392393394395396397398399400401402403404405406407408409410411412413414415416417418419420421422423424425426427428429430431432433434435436437438439440441442443444445446447448449450451452453454455456457458459460461462463464465466467468469470471472473474
  1. #include <torch/extension.h>
  2. #include "ATen/cuda/CUDAContext.h"
  3. #include <c10/cuda/CUDAGuard.h>
  4. #include "ln.h"
  5. /*
  6. Supported Type combinations:
  7. input residual compute weights output
  8. ============================================
  9. fp32 fp32 fp32 fp32 fp32
  10. fp16 fp32 fp32 fp32 fp16
  11. fp16 fp16 fp32 fp32 fp16
  12. bf16 fp32 fp32 fp32 bf16
  13. bf16 bf16 fp32 fp32 bf16
  14. fp16 fp16 fp32 fp16 fp16
  15. bf16 bf16 fp32 bf16 bf16
  16. Remarks:
  17. Output type = Input type
  18. Compute always in FP32
  19. */
  20. namespace layer_norm {
  21. // Create registries and provide runtime versions of config hash functions.
  22. FwdRegistry FWD_FUNCS;
  23. BwdRegistry BWD_FUNCS;
  24. ////////////////////////////////////////////////////////////////////////////////////////////////////
  25. uint32_t get_type_id(torch::Dtype dtype){
  26. if( dtype == torch::kFloat16 ) {
  27. return TypeId<fp16>::Value;
  28. } else if( dtype == torch::kBFloat16 ) {
  29. return TypeId<bf16>::Value;
  30. } else if( dtype == torch::kFloat32 ) {
  31. return TypeId<fp32>::Value;
  32. } else {
  33. TORCH_CHECK(false, "Type not supported: ", dtype);
  34. }
  35. }
  36. ////////////////////////////////////////////////////////////////////////////////////////////////////
  37. uint64_t get_key(torch::Dtype wtype, torch::Dtype itype, torch::Dtype rtype, torch::Dtype otype, torch::Dtype ctype, uint64_t hidden_size) {
  38. using namespace layer_norm;
  39. uint64_t type_key = get_type_id(wtype) | (get_type_id(itype) << 2) | (get_type_id(rtype) << 4) | (get_type_id(otype) << 6) | (get_type_id(ctype) << 8);
  40. uint64_t launcher_key = (type_key << 32) | hidden_size;
  41. return launcher_key;
  42. }
  43. } // namespace layer_norm
  44. ////////////////////////////////////////////////////////////////////////////////////////////////////
  45. layer_norm::FwdFunction & get_fwd_launcher(torch::Dtype wtype, torch::Dtype itype, torch::Dtype rtype, torch::Dtype otype, torch::Dtype ctype, uint32_t hidden_size) {
  46. auto iter = layer_norm::FWD_FUNCS.find(layer_norm::get_key(wtype, itype, rtype, otype, ctype, hidden_size));
  47. if( iter != layer_norm::FWD_FUNCS.end() ) {
  48. return iter->second;
  49. } else {
  50. TORCH_CHECK(false, "FWD: Unsupported hidden_size or types: ", hidden_size, wtype, itype, rtype, otype, ctype);
  51. }
  52. }
  53. ////////////////////////////////////////////////////////////////////////////////////////////////////
  54. layer_norm::BwdFunction & get_bwd_launcher(torch::Dtype wtype, torch::Dtype itype, torch::Dtype rtype, torch::Dtype otype, torch::Dtype ctype, uint32_t hidden_size) {
  55. auto iter = layer_norm::BWD_FUNCS.find(layer_norm::get_key(wtype, itype, rtype, otype, ctype, hidden_size));
  56. if( iter != layer_norm::BWD_FUNCS.end() ) {
  57. return iter->second;
  58. } else {
  59. TORCH_CHECK(false, "BWD: Unsupported hidden_size or types: ", hidden_size, wtype, itype, rtype, otype, ctype);
  60. }
  61. }
  62. ////////////////////////////////////////////////////////////////////////////////////////////////////
  63. std::vector<at::Tensor> dropout_add_ln_fwd(const at::Tensor &x0, // Input: BxSxhidden_size
  64. c10::optional<const at::Tensor> &residual_, // Residual: BxSxhidden_size
  65. const at::Tensor &gamma, // hidden_size
  66. c10::optional<const at::Tensor> &beta_, // hidden_size
  67. c10::optional<const at::Tensor> &rowscale_, // BxS
  68. c10::optional<const at::Tensor> &colscale_, // hidden_size
  69. c10::optional<const at::Tensor> &x0_subset_, // BxS
  70. c10::optional<const at::Tensor> &z_subset_, // BxS
  71. const float dropout_p,
  72. const float epsilon,
  73. const float rowscale_const,
  74. const int64_t z_numrows,
  75. c10::optional<at::Generator> gen_,
  76. bool residual_in_fp32=false,
  77. bool is_rms_norm=false
  78. ) {
  79. auto itype = x0.scalar_type();
  80. auto rtype = residual_.has_value()
  81. ? residual_.value().scalar_type()
  82. : (residual_in_fp32 ? torch::kFloat32 : x0.scalar_type());
  83. auto wtype = gamma.scalar_type();
  84. auto otype = itype;
  85. auto ctype = torch::kFloat32;
  86. auto mtype = torch::kUInt8;
  87. TORCH_CHECK(x0.is_cuda())
  88. TORCH_CHECK(gamma.is_cuda())
  89. TORCH_CHECK(x0.is_contiguous());
  90. // c10::IntArrayRef does not own the storage, so we need to construct a vector.
  91. // Otherwise just constructing IntArrayRef({blah}) will cause unintialized memory because
  92. // blah is then deallocated.
  93. std::vector<int64_t> sizes_vec {!x0_subset_.has_value() ? x0.size(0) : x0_subset_.value().size(0), x0.size(1)};
  94. auto sizes = c10::IntArrayRef(sizes_vec);
  95. TORCH_CHECK(x0.dim() == 2);
  96. TORCH_CHECK(sizes.size() == 2);
  97. const int rows = sizes[0];
  98. const int cols = sizes[1];
  99. auto hidden_size = gamma.numel();
  100. if (beta_.has_value()) {
  101. auto beta = beta_.value();
  102. TORCH_CHECK(beta.dtype() == wtype);
  103. TORCH_CHECK(beta.is_cuda())
  104. TORCH_CHECK(beta.is_contiguous());
  105. TORCH_CHECK(gamma.sizes() == beta.sizes());
  106. }
  107. if (residual_.has_value()) {
  108. auto residual = residual_.value();
  109. TORCH_CHECK(residual.is_cuda())
  110. TORCH_CHECK(residual.is_contiguous());
  111. TORCH_CHECK(residual.sizes() == sizes);
  112. }
  113. if (rowscale_.has_value()) {
  114. auto rowscale = rowscale_.value();
  115. TORCH_CHECK(rowscale.is_cuda())
  116. TORCH_CHECK(rowscale.is_contiguous());
  117. TORCH_CHECK(rowscale.sizes() == c10::IntArrayRef{rows});
  118. TORCH_CHECK(rowscale.dtype() == itype);
  119. }
  120. if (colscale_.has_value()) {
  121. auto colscale = colscale_.value();
  122. TORCH_CHECK(colscale.is_cuda())
  123. TORCH_CHECK(colscale.is_contiguous());
  124. TORCH_CHECK(colscale.sizes() == c10::IntArrayRef{cols});
  125. TORCH_CHECK(colscale.dtype() == wtype);
  126. }
  127. if (x0_subset_.has_value()) {
  128. auto x0_subset = x0_subset_.value();
  129. TORCH_CHECK(x0_subset.is_cuda())
  130. TORCH_CHECK(x0_subset.is_contiguous());
  131. TORCH_CHECK(x0_subset.sizes() == c10::IntArrayRef{rows});
  132. TORCH_CHECK(x0_subset.dtype() == torch::kInt32);
  133. TORCH_CHECK(z_subset_.has_value());
  134. auto z_subset = z_subset_.value();
  135. TORCH_CHECK(z_subset.is_cuda());
  136. TORCH_CHECK(z_subset.is_contiguous());
  137. TORCH_CHECK(z_subset.sizes() == c10::IntArrayRef{rows});
  138. TORCH_CHECK(z_subset.dtype() == torch::kInt32);
  139. }
  140. TORCH_CHECK(hidden_size == cols);
  141. TORCH_CHECK((hidden_size % 8 == 0) && (hidden_size <= 6144));
  142. TORCH_CHECK(epsilon >= 0.f);
  143. // Otherwise the kernel will be launched from cuda:0 device
  144. // Cast to char to avoid compiler warning about narrowing
  145. at::cuda::CUDAGuard device_guard{(char)x0.get_device()};
  146. auto opts = x0.options();
  147. bool save_x = residual_.has_value() || (dropout_p > 0.f) || rowscale_.has_value() || colscale_.has_value() || x0_subset_.has_value() || (itype != rtype);
  148. at::Tensor x;
  149. if (save_x) { x = torch::empty(sizes, opts.dtype(rtype)); }
  150. at::Tensor dmask;
  151. if (dropout_p > 0.f) { dmask = torch::empty(x0.sizes(), opts.dtype(mtype)); };
  152. auto z = torch::empty(z_subset_.has_value() ? c10::IntArrayRef{z_numrows, cols} : sizes, opts.dtype(otype));
  153. auto mu = torch::empty({ rows }, opts.dtype(ctype));
  154. auto rsigma = torch::empty({ rows }, opts.dtype(ctype));
  155. layer_norm::LaunchParams<layer_norm::FwdParams> launch_params;
  156. launch_params.props = at::cuda::getCurrentDeviceProperties();
  157. launch_params.stream = at::cuda::getCurrentCUDAStream().stream();
  158. TORCH_CHECK(dropout_p < 1.f);
  159. launch_params.params.dropout_keep_p = 1.f - dropout_p;
  160. launch_params.params.residual = residual_.has_value() ? residual_.value().data_ptr() : nullptr;
  161. launch_params.params.rowscale = rowscale_.has_value() ? rowscale_.value().data_ptr() : nullptr;
  162. launch_params.params.colscale = colscale_.has_value() ? colscale_.value().data_ptr() : nullptr;
  163. launch_params.params.x0_subset = x0_subset_.has_value() ? x0_subset_.value().data_ptr() : nullptr;
  164. launch_params.params.z_subset = z_subset_.has_value() ? z_subset_.value().data_ptr() : nullptr;
  165. auto gen = at::get_generator_or_default<at::CUDAGeneratorImpl>(
  166. gen_, at::cuda::detail::getDefaultCUDAGenerator());
  167. auto round_multiple = [](int x, int m) { return (x + m - 1) / m * m; };
  168. const int multiple = hidden_size <= 1536 ? 256 : (hidden_size <= 3072 ? 512 : 1024);
  169. // Request the kernel launcher.
  170. auto launcher = get_fwd_launcher(wtype, itype, rtype, otype, ctype, round_multiple(hidden_size, multiple));
  171. // Query the kernel-specific launch parameters.
  172. launcher(launch_params, true);
  173. at::Tensor workspace, barrier;
  174. // Set the kernel runtime parameters.
  175. layer_norm::FwdParams &params = launch_params.params;
  176. params.rows = rows;
  177. params.cols = cols;
  178. params.x0 = x0.data_ptr();
  179. params.x = save_x ? x.data_ptr() : nullptr;
  180. params.dmask = dropout_p > 0.f ? dmask.data_ptr() : nullptr;
  181. params.mu = mu.data_ptr();
  182. params.rs = rsigma.data_ptr();
  183. params.gamma = gamma.data_ptr();
  184. params.beta = beta_.has_value() ? beta_.value().data_ptr() : nullptr;
  185. params.z = z.data_ptr();
  186. params.epsilon = epsilon;
  187. params.dropout_scale = 1.f / (1.f - dropout_p);
  188. params.inverse_cols = 1.f / float(params.cols);
  189. params.rowscale_const = rowscale_const;
  190. params.is_rms_norm = is_rms_norm;
  191. if (dropout_p > 0.f) {
  192. // number of times random will be generated per thread, to offset philox counter in thc random
  193. // state
  194. int64_t counter_offset = launch_params.elts_per_thread;
  195. // See Note [Acquire lock when using random generators]
  196. {
  197. std::lock_guard<std::mutex> lock(gen->mutex_);
  198. params.philox_args = gen->philox_cuda_state(counter_offset);
  199. }
  200. }
  201. if( launch_params.barrier_size > 0 ) {
  202. auto options = x0.options();
  203. barrier = torch::zeros(launch_params.barrier_size, options.dtype(torch::kInt32));
  204. workspace = torch::empty(launch_params.workspace_bytes, options.dtype(torch::kChar));
  205. params.workspace = workspace.data_ptr();
  206. params.barrier = barrier.data_ptr<int>();
  207. }
  208. // Launch the kernel.
  209. launcher(launch_params, false);
  210. return { z, x, dmask, mu, rsigma };
  211. }
  212. ////////////////////////////////////////////////////////////////////////////////////////////////////
  213. std::vector<at::Tensor> dropout_add_ln_bwd(const at::Tensor &dz, // BxSxhidden_size
  214. c10::optional<const at::Tensor> &dx_, // BxSxhidden_size
  215. const at::Tensor &x, // BxSxhidden_size
  216. c10::optional<const at::Tensor> &x0_, // BxSxhidden_size
  217. c10::optional<const at::Tensor> &dmask_, // BxSxhidden_size
  218. const at::Tensor &mu, // BxS, FP32!
  219. const at::Tensor &rsigma, // BxS, FP32!
  220. const at::Tensor &gamma, // hidden_size
  221. c10::optional<const at::Tensor> &rowscale_, // BxS
  222. c10::optional<const at::Tensor> &colscale_, // hidden_size
  223. c10::optional<const at::Tensor> &x0_subset_, // BxS
  224. c10::optional<const at::Tensor> &z_subset_, // BxS
  225. const float dropout_p,
  226. const float rowscale_const,
  227. const int64_t x0_numrows,
  228. const bool has_residual,
  229. bool is_rms_norm=false
  230. ) {
  231. auto itype = dz.scalar_type();
  232. auto rtype = x.scalar_type();
  233. auto wtype = gamma.scalar_type();
  234. auto otype = itype;
  235. auto ctype = torch::kFloat32;
  236. auto mtype = torch::kUInt8;
  237. if (dropout_p > 0.f) { TORCH_CHECK(dmask_.has_value()); }
  238. TORCH_CHECK(dz.dtype() == otype);
  239. TORCH_CHECK(mu.dtype() == ctype);
  240. TORCH_CHECK(rsigma.dtype() == ctype);
  241. TORCH_CHECK(x.is_cuda());
  242. TORCH_CHECK(dz.is_cuda());
  243. TORCH_CHECK(mu.is_cuda());
  244. TORCH_CHECK(rsigma.is_cuda());
  245. TORCH_CHECK(gamma.is_cuda());
  246. TORCH_CHECK(x.is_contiguous());
  247. TORCH_CHECK(dz.is_contiguous());
  248. auto sizes = x.sizes();
  249. TORCH_CHECK(sizes.size() == 2);
  250. auto rows = sizes[0];
  251. auto cols = sizes[1];
  252. TORCH_CHECK(dz.dim() == 2);
  253. TORCH_CHECK(dz.size(1) == cols);
  254. // c10::IntArrayRef does not own the storage, so we need to construct a vector.
  255. // Otherwise just constructing IntArrayRef({blah}) will cause unintialized memory because
  256. // blah is then deallocated.
  257. std::vector<int64_t> x0_sizes_vec {!x0_subset_.has_value() ? rows : x0_numrows, cols};
  258. auto x0_sizes = c10::IntArrayRef(x0_sizes_vec);
  259. if (dx_.has_value()) {
  260. auto dx = dx_.value();
  261. TORCH_CHECK(dx.dtype() == rtype);
  262. TORCH_CHECK(dx.is_cuda())
  263. TORCH_CHECK(dx.is_contiguous());
  264. TORCH_CHECK(dx.sizes() == sizes);
  265. }
  266. if (dmask_.has_value()) {
  267. auto dmask = dmask_.value();
  268. TORCH_CHECK(dmask.dtype() == mtype);
  269. TORCH_CHECK(dmask.is_cuda());
  270. TORCH_CHECK(dmask.is_contiguous());
  271. TORCH_CHECK(dmask.sizes() == x0_sizes);
  272. }
  273. if (rowscale_.has_value()) {
  274. auto rowscale = rowscale_.value();
  275. TORCH_CHECK(rowscale.is_cuda())
  276. TORCH_CHECK(rowscale.is_contiguous());
  277. TORCH_CHECK(rowscale.sizes() == c10::IntArrayRef{rows});
  278. TORCH_CHECK(rowscale.dtype() == itype);
  279. }
  280. if (colscale_.has_value()) {
  281. auto colscale = colscale_.value();
  282. TORCH_CHECK(colscale.is_cuda())
  283. TORCH_CHECK(colscale.is_contiguous());
  284. TORCH_CHECK(colscale.sizes() == c10::IntArrayRef{cols});
  285. TORCH_CHECK(colscale.dtype() == wtype);
  286. TORCH_CHECK(x0_.has_value());
  287. auto x0 = x0_.value();
  288. TORCH_CHECK(x0.is_cuda())
  289. TORCH_CHECK(x0.is_contiguous());
  290. TORCH_CHECK(x0.sizes() == x0_sizes);
  291. TORCH_CHECK(x0.dtype() == itype);
  292. }
  293. if (x0_subset_.has_value()) {
  294. auto x0_subset = x0_subset_.value();
  295. TORCH_CHECK(x0_subset.is_cuda())
  296. TORCH_CHECK(x0_subset.is_contiguous());
  297. TORCH_CHECK(x0_subset.sizes() == c10::IntArrayRef{rows});
  298. TORCH_CHECK(x0_subset.dtype() == torch::kInt32);
  299. TORCH_CHECK(z_subset_.has_value());
  300. auto z_subset = z_subset_.value();
  301. TORCH_CHECK(z_subset.is_cuda());
  302. TORCH_CHECK(z_subset.is_contiguous());
  303. TORCH_CHECK(z_subset.sizes() == c10::IntArrayRef{rows});
  304. TORCH_CHECK(z_subset.dtype() == torch::kInt32);
  305. }
  306. auto hidden_size = gamma.numel();
  307. TORCH_CHECK(hidden_size == cols);
  308. TORCH_CHECK((hidden_size % 8 == 0) && (hidden_size <= 6144));
  309. TORCH_CHECK(mu.numel() == rows);
  310. TORCH_CHECK(mu.sizes() == rsigma.sizes());
  311. TORCH_CHECK(gamma.numel() == cols);
  312. // Otherwise the kernel will be launched from cuda:0 device
  313. // Cast to char to avoid compiler warning about narrowing
  314. at::cuda::CUDAGuard device_guard{(char)dz.get_device()};
  315. auto opts = x.options();
  316. auto dx0 = torch::empty(x0_sizes, opts.dtype(itype));
  317. at::Tensor dresidual;
  318. if (has_residual) { dresidual = torch::empty_like(x, opts.dtype(rtype)); }
  319. auto dgamma = torch::empty_like(gamma);
  320. auto dbeta = torch::empty_like(gamma);
  321. at::Tensor dcolscale;
  322. if (colscale_.has_value()) {
  323. dcolscale = torch::empty_like(colscale_.value());
  324. }
  325. layer_norm::LaunchParams<layer_norm::BwdParams> launch_params;
  326. launch_params.stream = at::cuda::getCurrentCUDAStream().stream();
  327. launch_params.props = at::cuda::getCurrentDeviceProperties();
  328. TORCH_CHECK(dropout_p < 1.f);
  329. launch_params.params.dropout_keep_p = 1.f - dropout_p;
  330. launch_params.params.dresidual = has_residual ? dresidual.data_ptr() : nullptr;
  331. launch_params.params.rowscale = rowscale_.has_value() ? rowscale_.value().data_ptr() : nullptr;
  332. launch_params.params.colscale = colscale_.has_value() ? colscale_.value().data_ptr() : nullptr;
  333. launch_params.params.x0_subset = x0_subset_.has_value() ? x0_subset_.value().data_ptr() : nullptr;
  334. launch_params.params.z_subset = z_subset_.has_value() ? z_subset_.value().data_ptr() : nullptr;
  335. auto round_multiple = [](int x, int m) { return (x + m - 1) / m * m; };
  336. const int multiple = hidden_size <= 1536 ? 256 : (hidden_size <= 3072 ? 512 : 1024);
  337. auto launcher = get_bwd_launcher(wtype, itype, rtype, otype, ctype, round_multiple(hidden_size, multiple));
  338. launcher(launch_params, true);
  339. auto dgamma_part = torch::empty({ launch_params.params.ctas_per_col, hidden_size }, opts.dtype(ctype));
  340. auto dbeta_part = torch::empty({ launch_params.params.ctas_per_col, hidden_size }, opts.dtype(ctype));
  341. at::Tensor dcolscale_part;
  342. if (colscale_.has_value()) {
  343. dcolscale_part = torch::empty({ launch_params.params.ctas_per_col, hidden_size }, opts.dtype(ctype));
  344. }
  345. at::Tensor workspace, barrier;
  346. layer_norm::BwdParams &params = launch_params.params;
  347. params.rows = rows;
  348. params.cols = cols;
  349. params.x = x.data_ptr();
  350. params.x0 = x0_.has_value() ? x0_.value().data_ptr() : nullptr;
  351. params.dmask = dropout_p > 0.f ? dmask_.value().data_ptr() : nullptr;
  352. params.mu = mu.data_ptr();
  353. params.rs = rsigma.data_ptr();
  354. params.gamma = gamma.data_ptr();
  355. params.dz = dz.data_ptr();
  356. params.dx = dx_.has_value() ? dx_.value().data_ptr() : nullptr;
  357. params.dx0 = dx0.data_ptr();
  358. params.dbeta = dbeta.data_ptr();
  359. params.dgamma = dgamma.data_ptr();
  360. params.dcolscale = colscale_.has_value() ? dcolscale.data_ptr() : nullptr;
  361. params.dbeta_part = dbeta_part.data_ptr();
  362. params.dgamma_part = dgamma_part.data_ptr();
  363. params.dcolscale_part = colscale_.has_value() ? dcolscale_part.data_ptr() : nullptr;
  364. params.dropout_scale = 1.f / (1.f - dropout_p);
  365. params.inverse_cols = 1.f / float(params.cols);
  366. params.rowscale_const = rowscale_const;
  367. params.is_rms_norm = is_rms_norm;
  368. if( launch_params.barrier_size > 0 ) {
  369. // TODO Any way to avoid this?
  370. barrier = torch::zeros(launch_params.barrier_size, opts.dtype(torch::kInt32));
  371. workspace = torch::empty(launch_params.workspace_bytes, opts.dtype(torch::kChar));
  372. params.workspace = workspace.data_ptr();
  373. params.barrier = barrier.data_ptr<int>();
  374. }
  375. launcher(launch_params, false);
  376. std::vector<at::Tensor> result = { dx0, dresidual, dgamma, dbeta, dgamma_part, dbeta_part };
  377. if (colscale_.has_value()) {
  378. result.push_back(dcolscale);
  379. result.push_back(dcolscale_part);
  380. }
  381. return result;
  382. }
  383. ////////////////////////////////////////////////////////////////////////////////////////////////////
  384. PYBIND11_MODULE(TORCH_EXTENSION_NAME, m) {
  385. m.doc() = "CUDA DropoutAddLayerNorm";
  386. m.def("dropout_add_ln_fwd", &dropout_add_ln_fwd, "Run Dropout + Add + LayerNorm forward kernel",
  387. py::arg("x0"), py::arg("residual"), py::arg("gamma"), py::arg("beta"),
  388. py::arg("rowscale_"), py::arg("colscale_"), py::arg("x0_subset_"), py::arg("z_subset_"),
  389. py::arg("dropout_p"), py::arg("epsilon"), py::arg("rowscale_const"), py::arg("z_numrows"),
  390. py::arg("gen_"), py::arg("residual_in_fp32")=false, py::arg("is_rms_norm")=false);
  391. m.def("dropout_add_ln_bwd", &dropout_add_ln_bwd, "Run Dropout + Add + LayerNorm backward kernel",
  392. py::arg("dz"), py::arg("dx_"), py::arg("x"), py::arg("x0_"), py::arg("dmask_"), py::arg("mu"),
  393. py::arg("rsigma"), py::arg("gamma"), py::arg("rowscale_"), py::arg("colscale_"),
  394. py::arg("x0_subset_"), py::arg("z_subset_"), py::arg("dropout_p"), py::arg("rowscale_const"),
  395. py::arg("x0_numrows"), py::arg("has_residual"), py::arg("is_rms_norm")=false);
  396. }