ufunc.py 17 KB

123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121122123124125126127128129130131132133134135136137138139140141142143144145146147148149150151152153154155156157158159160161162163164165166167168169170171172173174175176177178179180181182183184185186187188189190191192193194195196197198199200201202203204205206207208209210211212213214215216217218219220221222223224225226227228229230231232233234235236237238239240241242243244245246247248249250251252253254255256257258259260261262263264265266267268269270271272273274275276277278279280281282283284285286287288289290291292293294295296297298299300301302303304305306307308309310311312313314315316317318319320321322323324325326327328329330331332333334335336337338339340341342343344345346347348349350351352353354355356357358359360361362363364365366367368369370371372373374375376377378379380381382383384385386387388389390391392393394395396397398399400401402403404405406407408409410411412413414415416417418419420421422423424425426427428429430431432433434435436437438439440441442443444445446447448449450451452453454455456457458459460461462463464465466467468469470471472473474475476477478479480481482483484485486487488489490491492493494495496497498499500501502503504505506507508509510511512513514515516517518519520521522523524525526527528529530531532533534535536537538539540541542543544545
  1. from dataclasses import dataclass
  2. from typing import Dict, List, Optional, Sequence, Tuple, Union
  3. import torchgen.api.ufunc as ufunc
  4. from torchgen.api.translate import translate
  5. from torchgen.api.types import (
  6. BaseCType,
  7. Binding,
  8. CType,
  9. Expr,
  10. NamedCType,
  11. opmath_t,
  12. scalar_t,
  13. StructuredImplSignature,
  14. VectorizedCType,
  15. )
  16. from torchgen.api.ufunc import UfunctorBindings
  17. from torchgen.context import with_native_function
  18. from torchgen.model import (
  19. Argument,
  20. BaseTy,
  21. BaseType,
  22. DispatchKey,
  23. NativeFunctionsGroup,
  24. ScalarType,
  25. UfuncKey,
  26. )
  27. from torchgen.utils import OrderedSet
  28. # ~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~ #
  29. #
  30. # CUDA STUFF
  31. #
  32. # ~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~ #
  33. # NB: not bothering to generate dispatch stub forward declaration in header,
  34. # we can just paste it whereever necessary
  35. # TODO: use BackendIndex
  36. # dispatch_key: DispatchKey # only CPU/CUDA right now
  37. # Represents functors for implementing CUDA ufuncs.
  38. # Functors are templated by scalar_t because when USERS instantiate functors
  39. # they are templated. A functor looks something like this:
  40. #
  41. # template <typename scalar_t>
  42. # struct CUDAFunctorOnSelf_add {
  43. # using opmath_t = at::opmath_type<scalar_t>;
  44. # opmath_t other_;
  45. # opmath_t alpha_;
  46. # CUDAFunctorOnSelf_add(opmath_t other, opmath_t alpha)
  47. # : other_(other), alpha_(alpha) {}
  48. # __device__ scalar_t operator()(scalar_t self) {
  49. # return ufunc::add(static_cast<opmath_t>(self), other_, alpha_);
  50. # }
  51. # };
  52. #
  53. @dataclass(frozen=True)
  54. class UfunctorSignature:
  55. g: NativeFunctionsGroup
  56. scalar_tensor_idx: Optional[int]
  57. name: str
  58. def arguments(self) -> UfunctorBindings:
  59. return ufunc.ufunctor_arguments(
  60. self.g, scalar_tensor_idx=self.scalar_tensor_idx, scalar_t=scalar_t
  61. )
  62. def fields(self) -> List[Binding]:
  63. # fields are renamed to have a trailing underscore, as is conventional
  64. return [b.rename(f"{b.name}_") for b in self.arguments().ctor]
  65. def returns_type(self) -> CType:
  66. # TODO: don't hardcode; return type will be inferred based on tags on
  67. # the native function
  68. return BaseCType(scalar_t)
  69. def decl_fields(self) -> str:
  70. return "\n".join(f"{f.type} {f.name};" for f in self.fields())
  71. def inline_defn_ctor(self) -> str:
  72. args_str = ", ".join(a.decl() for a in self.arguments().ctor)
  73. # NB: hypothetically could do this with translate but the
  74. # transition here is very regular
  75. init_str = ", ".join(f"{a.name}_({a.name})" for a in self.arguments().ctor)
  76. return f"{self.name}({args_str}) : {init_str} {{}}"
  77. def decl_apply(self) -> str:
  78. args_str = ", ".join(a.decl() for a in self.arguments().apply)
  79. return f"{self.returns_type().cpp_type()} operator()({args_str}) const"
  80. @dataclass(frozen=True)
  81. class UfuncSignature:
  82. g: NativeFunctionsGroup
  83. name: str
  84. compute_t: CType
  85. def arguments(self) -> List[Binding]:
  86. return ufunc.ufunc_arguments(self.g, compute_t=self.compute_t)
  87. def call(self, ctx: Sequence[Union[Binding, Expr]]) -> str:
  88. return f"{self.name}({', '.join(a.expr for a in translate(ctx, self.arguments()))})"
  89. # steps:
  90. # 1. take the functional signature
  91. # 2. use api.ufunc to convert it to template signature. this establishes
  92. # the type of the template function
  93. # 3. use api.ufunc (II) to generate a split struct / operator() signature.
  94. # this establish context in which we call the template signature
  95. #
  96. # StructuredImplSignature context
  97. # ~> functor constructor sig
  98. #
  99. # Functor constructor context
  100. # ~> functor fields sig
  101. #
  102. # Functor apply context (functor fields + functor apply sig)
  103. # ~> template sig
  104. #
  105. def eligible_for_binary_scalar_specialization(g: NativeFunctionsGroup) -> bool:
  106. num_tensors = sum(
  107. 1 for a in g.functional.func.arguments.flat_non_out if a.type.is_tensor_like()
  108. )
  109. return num_tensors == 2
  110. def compute_ufunc_cuda_functors(
  111. g: NativeFunctionsGroup,
  112. ) -> Tuple[Dict[ScalarType, Dict[UfuncKey, UfunctorSignature]], str]:
  113. # First, build the functors.
  114. ufunctor_sigs: Dict[ScalarType, Dict[UfuncKey, UfunctorSignature]] = {}
  115. ufunctors: List[str] = []
  116. loops = g.out.ufunc_inner_loop
  117. scalar_tensor_idx_lookup = {
  118. UfuncKey.CUDAFunctorOnSelf: 1,
  119. UfuncKey.CUDAFunctorOnOther: 0,
  120. UfuncKey.CUDAFunctor: None,
  121. }
  122. if eligible_for_binary_scalar_specialization(g):
  123. keys = [
  124. UfuncKey.CUDAFunctorOnSelf,
  125. UfuncKey.CUDAFunctorOnOther,
  126. UfuncKey.CUDAFunctor,
  127. ]
  128. else:
  129. keys = [UfuncKey.CUDAFunctor]
  130. for k in [UfuncKey.CUDAFunctorOnSelf, UfuncKey.CUDAFunctorOnOther]:
  131. assert k not in loops, f"cannot use {k} on non-binary function"
  132. for k in keys:
  133. # If the key was directly defined, skip functor codegen; we assume the
  134. # user already done it for us
  135. if k in loops:
  136. ufunctor_sig = UfunctorSignature(
  137. g, scalar_tensor_idx=scalar_tensor_idx_lookup[k], name=loops[k].name
  138. )
  139. for dtype in loops[k].supported_dtypes:
  140. ufunctor_sigs.setdefault(dtype, {})[k] = ufunctor_sig
  141. continue
  142. # Note [ScalarOnly and Generic must match names for CUDA]
  143. # ~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~
  144. # Otherwise, look in ANY of the generic entries. For simplicity of
  145. # codegen, both ScalarOnly and Generic are defined, the ufunc name
  146. # must match (if they didn't match, we'd have to generate distinct
  147. # functors per dtype, which is awful, so we're not going to do it unless
  148. # someone really forces us to)
  149. ufunc_name = None
  150. supported_dtypes: OrderedSet[ScalarType] = OrderedSet()
  151. for lk in [UfuncKey.ScalarOnly, UfuncKey.Generic]:
  152. if lk not in loops:
  153. continue
  154. if ufunc_name is None:
  155. ufunc_name = loops[lk].name
  156. else:
  157. # See Note [ScalarOnly and Generic must match names for CUDA]
  158. assert (
  159. ufunc_name == loops[lk].name
  160. ), "ScalarOnly and Generic must have same ufunc name"
  161. supported_dtypes |= loops[lk].supported_dtypes
  162. assert ufunc_name is not None
  163. name = f"{k}_{ufunc_name}"
  164. ufunctor_sig = UfunctorSignature(
  165. g, scalar_tensor_idx=scalar_tensor_idx_lookup[k], name=name
  166. )
  167. for dtype in supported_dtypes:
  168. ufunctor_sigs.setdefault(dtype, {})[k] = ufunctor_sig
  169. ufunc_sig = UfuncSignature(
  170. g, name=f"ufunc::{ufunc_name}", compute_t=BaseCType(opmath_t)
  171. )
  172. apply_ctx = ufunctor_sig.fields() + ufunctor_sig.arguments().apply
  173. ufunctors.append(
  174. f"""
  175. template <typename scalar_t>
  176. struct {ufunctor_sig.name} {{
  177. using opmath_t = at::opmath_type<scalar_t>;
  178. {ufunctor_sig.decl_fields()}
  179. {ufunctor_sig.inline_defn_ctor()}
  180. __device__ {ufunctor_sig.decl_apply()} {{
  181. return {ufunc_sig.call(apply_ctx)};
  182. }}
  183. }};
  184. """
  185. )
  186. return ufunctor_sigs, "\n".join(ufunctors)
  187. @dataclass(frozen=True)
  188. class BinaryScalarSpecializationConfig:
  189. scalar_idx: int
  190. ctor_tensor: str
  191. ufunc_key: UfuncKey
  192. BinaryScalarSpecializationConfigs = [
  193. BinaryScalarSpecializationConfig(
  194. scalar_idx=0,
  195. ctor_tensor="self",
  196. ufunc_key=UfuncKey.CUDAFunctorOnOther,
  197. ),
  198. BinaryScalarSpecializationConfig(
  199. scalar_idx=1,
  200. ctor_tensor="other",
  201. ufunc_key=UfuncKey.CUDAFunctorOnSelf,
  202. ),
  203. ]
  204. def compute_ufunc_cuda_dtype_body(
  205. g: NativeFunctionsGroup,
  206. dtype: ScalarType,
  207. inner_loops: Dict[UfuncKey, UfunctorSignature],
  208. parent_ctx: Sequence[Binding],
  209. ) -> str:
  210. body = "using opmath_t = at::opmath_type<scalar_t>;"
  211. body += "if (false) {}\n" # for ease of codegen
  212. for config in BinaryScalarSpecializationConfigs:
  213. if config.ufunc_key not in inner_loops:
  214. continue
  215. ufunctor_sig = inner_loops[config.ufunc_key]
  216. scalar_idx = config.scalar_idx + 1
  217. # Make a copy and at the same time widen the type (not permissible
  218. # without copy; we don't want to mutate the input argument anyway)
  219. ctx: List[Union[Expr, Binding]] = list(parent_ctx)
  220. ctx.append(
  221. Expr(
  222. expr=f"iter.scalar_value<opmath_t>({scalar_idx})",
  223. type=NamedCType(config.ctor_tensor, BaseCType(opmath_t)),
  224. )
  225. )
  226. ufunctor_ctor_exprs_str = ", ".join(
  227. a.expr for a in translate(ctx, ufunctor_sig.arguments().ctor)
  228. )
  229. # NB: ufunctor must be allocated before iter.remove_operand is called,
  230. # as it relies on iter
  231. body += f"""\
  232. else if (iter.is_cpu_scalar({scalar_idx})) {{
  233. {ufunctor_sig.name}<scalar_t> ufunctor({ufunctor_ctor_exprs_str});
  234. iter.remove_operand({scalar_idx});
  235. gpu_kernel(iter, ufunctor);
  236. }}"""
  237. ufunctor_sig = inner_loops[UfuncKey.CUDAFunctor]
  238. ufunctor_ctor_exprs_str = ", ".join(
  239. a.expr for a in translate(parent_ctx, ufunctor_sig.arguments().ctor)
  240. )
  241. body += f"""
  242. else {{
  243. gpu_kernel(iter, {ufunctor_sig.name}<scalar_t>({ufunctor_ctor_exprs_str}));
  244. }}
  245. """
  246. return body
  247. @with_native_function
  248. def compute_ufunc_cuda(g: NativeFunctionsGroup) -> str:
  249. # First, build the functors, indexing them by dtype
  250. ufunctor_sigs, ufunctors = compute_ufunc_cuda_functors(g)
  251. # Next, build the conditionals
  252. sig = StructuredImplSignature(g, ufunc.kernel_name(g, DispatchKey.CUDA))
  253. dtype_cases = []
  254. for dtype, inner_ufunc_sigs in ufunctor_sigs.items():
  255. dtype_cases.append(
  256. f"""
  257. AT_DISPATCH_CASE(at::ScalarType::{dtype},
  258. [&]() {{
  259. {compute_ufunc_cuda_dtype_body(g, dtype, inner_ufunc_sigs, sig.arguments())}
  260. }}
  261. )
  262. """
  263. )
  264. dtype_cases_str = "\n".join(dtype_cases)
  265. stub_sig = StubSignature(g)
  266. return f"""
  267. {ufunctors}
  268. {stub_sig.type_defn()};
  269. {stub_sig.dispatch_decl()};
  270. {stub_sig.kernel_defn()} {{
  271. AT_DISPATCH_SWITCH(iter.common_dtype(), "{sig.name}",
  272. {dtype_cases_str}
  273. );
  274. }}
  275. REGISTER_DISPATCH({stub_sig.name}, &{stub_sig.kernel_name});
  276. {sig.defn()} {{
  277. {stub_sig.direct_call(sig.arguments())};
  278. }}
  279. """
  280. # ~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~ #
  281. #
  282. # CPU STUFF
  283. #
  284. # ~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~~ #
  285. @dataclass(frozen=True)
  286. class StubSignature:
  287. g: NativeFunctionsGroup
  288. @property
  289. def name(self) -> str:
  290. return f"{str(self.g.functional.func.name.name)}_stub"
  291. @property
  292. def kernel_name(self) -> str:
  293. return f"{str(self.g.functional.func.name.name)}_kernel"
  294. @property
  295. def type_name(self) -> str:
  296. return f"{str(self.g.functional.func.name.name)}_fn"
  297. def arguments(self) -> List[Binding]:
  298. return ufunc.stub_arguments(self.g)
  299. def type(self) -> str:
  300. cpp_args = self.arguments()
  301. return f"void(*)(TensorIteratorBase&, {', '.join(a.type for a in cpp_args)})"
  302. def dispatch_decl(self) -> str:
  303. return f"DECLARE_DISPATCH({self.type_name}, {self.name})"
  304. def dispatch_defn(self) -> str:
  305. return f"DEFINE_DISPATCH({self.name})"
  306. def kernel_defn(self) -> str:
  307. return f"void {self.kernel_name}(TensorIteratorBase& iter, {', '.join(a.defn() for a in self.arguments())})"
  308. def type_defn(self) -> str:
  309. return f"using {self.type_name} = {self.type()}"
  310. # must be called from context where this is TensorIteratorBase*
  311. def call(self, ctx: Sequence[Binding]) -> str:
  312. return f"{self.name}(device_type(), *this, {', '.join(a.expr for a in translate(ctx, self.arguments()))})"
  313. # used in CUDA to skip the unnecessary dynamic dispatch
  314. def direct_call(self, ctx: Sequence[Binding]) -> str:
  315. return f"{self.kernel_name}(*this, {', '.join(a.expr for a in translate(ctx, self.arguments()))})"
  316. @with_native_function
  317. def compute_ufunc_cpu(g: NativeFunctionsGroup) -> str:
  318. stub_sig = StubSignature(g)
  319. sig = StructuredImplSignature(g, ufunc.kernel_name(g, DispatchKey.CPU))
  320. return f"""
  321. {stub_sig.type_defn()};
  322. {stub_sig.dispatch_decl()};
  323. {stub_sig.dispatch_defn()};
  324. {sig.defn()} {{
  325. {stub_sig.call(sig.arguments())};
  326. }}
  327. """
  328. def compute_ufunc_cpu_dtype_body(
  329. g: NativeFunctionsGroup,
  330. dtype: ScalarType,
  331. inner_loops: Dict[UfuncKey, UfuncSignature],
  332. parent_ctx: Sequence[Binding],
  333. ) -> str:
  334. assert UfuncKey.CPUScalar in inner_loops, f"{dtype}, {inner_loops.keys()}"
  335. assert inner_loops.keys() <= {UfuncKey.CPUScalar, UfuncKey.CPUVector}
  336. scalar_loop = inner_loops[UfuncKey.CPUScalar]
  337. vec_loop = None
  338. if UfuncKey.CPUVector in inner_loops:
  339. vec_loop = inner_loops[UfuncKey.CPUVector]
  340. # NB: We DON'T use translate here, because translate is
  341. # incapable of CSE'ing the scalar accesses in case it is also
  342. # used by Vectorized; also, the unpacking here is very simple
  343. # and only affects Scalar; everything else is implicitly captured
  344. # by the lambda
  345. # Setup scalar in scope
  346. body = []
  347. ctx = []
  348. for b in parent_ctx:
  349. if isinstance(b.argument, Argument) and b.argument.type != BaseType(
  350. BaseTy.Scalar
  351. ):
  352. continue
  353. body.append(f"auto _s_{b.name} = {b.name}.to<scalar_t>();")
  354. ctx.append(Expr(f"_s_{b.name}", NamedCType(b.nctype.name, BaseCType(scalar_t))))
  355. if vec_loop is not None:
  356. for b in parent_ctx:
  357. if isinstance(b.argument, Argument) and b.argument.type != BaseType(
  358. BaseTy.Scalar
  359. ):
  360. continue
  361. body.append(
  362. f"auto _v_{b.name} = at::vec::Vectorized<scalar_t>(_s_{b.name});"
  363. )
  364. ctx.append(
  365. Expr(
  366. f"_v_{b.name}",
  367. NamedCType(b.nctype.name, VectorizedCType(BaseCType(scalar_t))),
  368. )
  369. )
  370. # Setup lambda signature
  371. # NB: simplified version of ufunctor_arguments
  372. scalar_bindings = []
  373. vec_bindings = []
  374. for a in g.functional.func.arguments.flat_non_out:
  375. if not a.type.is_tensor_like():
  376. continue
  377. assert a.type == BaseType(BaseTy.Tensor)
  378. scalar_bindings.append(
  379. Binding(
  380. name=a.name,
  381. nctype=NamedCType(a.name, BaseCType(scalar_t)),
  382. argument=a,
  383. )
  384. )
  385. if vec_loop is not None:
  386. vec_bindings.append(
  387. Binding(
  388. name=a.name,
  389. nctype=NamedCType(a.name, VectorizedCType(BaseCType(scalar_t))),
  390. argument=a,
  391. )
  392. )
  393. def with_ctx(b: Sequence[Binding]) -> List[Union[Expr, Binding]]:
  394. r: List[Union[Expr, Binding]] = []
  395. r.extend(ctx)
  396. r.extend(b)
  397. return r
  398. body_str = "\n".join(body)
  399. if vec_loop is not None:
  400. return f"""
  401. {body_str}
  402. cpu_kernel_vec(iter,
  403. [=]({', '.join(b.decl() for b in scalar_bindings)}) {{ return {scalar_loop.call(with_ctx(scalar_bindings))}; }},
  404. [=]({', '.join(b.decl() for b in vec_bindings)}) {{ return {vec_loop.call(with_ctx(vec_bindings))}; }}
  405. );
  406. """
  407. else:
  408. return f"""
  409. {body_str}
  410. cpu_kernel(iter,
  411. [=]({', '.join(b.decl() for b in scalar_bindings)}) {{ return {scalar_loop.call(with_ctx(scalar_bindings))}; }}
  412. );
  413. """
  414. @with_native_function
  415. def compute_ufunc_cpu_kernel(g: NativeFunctionsGroup) -> str:
  416. stub_sig = StubSignature(g)
  417. # Reindex the ufunc by dtypes; processing generic/scalaronly as well
  418. loops = g.out.ufunc_inner_loop
  419. ufunc_sigs: Dict[ScalarType, Dict[UfuncKey, UfuncSignature]] = {}
  420. for k in [UfuncKey.CPUScalar, UfuncKey.CPUVector]:
  421. lks = []
  422. # ORDER MATTERS: this specifies overriding precedence
  423. if k in loops: # should happen rarely
  424. lks.append(k)
  425. if UfuncKey.ScalarOnly in loops and k is UfuncKey.CPUScalar:
  426. lks.append(UfuncKey.ScalarOnly)
  427. if UfuncKey.Generic in loops:
  428. lks.append(UfuncKey.Generic)
  429. # TODO: don't hardcode ufunc:: namespace here, should be centralized smh
  430. for lk in lks:
  431. for dtype in loops[lk].supported_dtypes:
  432. compute_t: CType
  433. if k is UfuncKey.CPUScalar:
  434. compute_t = BaseCType(scalar_t)
  435. elif k is UfuncKey.CPUVector:
  436. compute_t = VectorizedCType(BaseCType(scalar_t))
  437. else:
  438. raise AssertionError()
  439. inner_ufunc_sigs = ufunc_sigs.setdefault(dtype, {})
  440. if k not in inner_ufunc_sigs:
  441. inner_ufunc_sigs[k] = UfuncSignature(
  442. g, name=f"ufunc::{loops[lk].name}", compute_t=compute_t
  443. )
  444. # Build the conditionals
  445. dtype_cases = []
  446. for dtype, inner_ufunc_sigs in ufunc_sigs.items():
  447. dtype_cases.append(
  448. f"""
  449. AT_DISPATCH_CASE(at::ScalarType::{dtype},
  450. [&]() {{
  451. {compute_ufunc_cpu_dtype_body(g, dtype, inner_ufunc_sigs, stub_sig.arguments())}
  452. }}
  453. )
  454. """
  455. )
  456. dtype_cases_str = "\n".join(dtype_cases)
  457. return f"""
  458. namespace {{
  459. {stub_sig.kernel_defn()} {{
  460. AT_DISPATCH_SWITCH(iter.common_dtype(), "{stub_sig.name}",
  461. {dtype_cases_str}
  462. );
  463. }}
  464. }} // anonymous namespace
  465. {stub_sig.type_defn()};
  466. {stub_sig.dispatch_decl()};
  467. REGISTER_DISPATCH({stub_sig.name}, &{stub_sig.kernel_name});
  468. """