Skip to content

Commit d61986c

Browse files
Add Decompostion for Aten_SafeSoftmaxOp (#3708)
Co-authored-by: Vivek Khandelwal <[email protected]>
1 parent edf725e commit d61986c

File tree

8 files changed

+160
-8
lines changed

8 files changed

+160
-8
lines changed

include/torch-mlir/Dialect/Torch/IR/GeneratedTorchOps.td

Lines changed: 25 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -8370,6 +8370,31 @@ def Torch_Aten_SoftmaxOp : Torch_Op<"aten._softmax", [
83708370
}];
83718371
}
83728372

8373+
def Torch_Aten_SafeSoftmaxOp : Torch_Op<"aten._safe_softmax", [
8374+
AllowsTypeRefinement,
8375+
HasValueSemantics,
8376+
ReadOnly
8377+
]> {
8378+
let summary = "Generated op for `aten::_safe_softmax : (Tensor, int, int?) -> (Tensor)`";
8379+
let arguments = (ins
8380+
AnyTorchTensorType:$self,
8381+
Torch_IntType:$dim,
8382+
AnyTorchOptionalIntType:$dtype
8383+
);
8384+
let results = (outs
8385+
AnyTorchOptionalTensorType:$result
8386+
);
8387+
let hasCustomAssemblyFormat = 1;
8388+
let extraClassDefinition = [{
8389+
ParseResult Aten_SafeSoftmaxOp::parse(OpAsmParser &parser, OperationState &result) {
8390+
return parseDefaultTorchOp(parser, result, 3, 1);
8391+
}
8392+
void Aten_SafeSoftmaxOp::print(OpAsmPrinter &printer) {
8393+
printDefaultTorchOp(printer, *this, 3, 1);
8394+
}
8395+
}];
8396+
}
8397+
83738398
def Torch_AtenMeanOp : Torch_Op<"aten.mean", [
83748399
AllowsTypeRefinement,
83758400
HasValueSemantics,

lib/Dialect/Torch/Transforms/AbstractInterpLibrary.cpp

Lines changed: 16 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -6772,6 +6772,10 @@ StringRef mlir::torch::Torch::getAbstractInterpLibrary() {
67726772
" %0 = call @__torch__.torch.jit._shape_functions.unary(%arg0) : (!torch.list<int>) -> !torch.list<int>\n"
67736773
" return %0 : !torch.list<int>\n"
67746774
" }\n"
6775+
" func.func @\"__torch_mlir_shape_fn.aten._safe_softmax\"(%arg0: !torch.list<int>, %arg1: !torch.int, %arg2: !torch.optional<int>) -> !torch.list<int> {\n"
6776+
" %0 = call @__torch__.torch.jit._shape_functions.unary(%arg0) : (!torch.list<int>) -> !torch.list<int>\n"
6777+
" return %0 : !torch.list<int>\n"
6778+
" }\n"
67756779
" func.func @\"__torch_mlir_shape_fn.aten.softmax.int\"(%arg0: !torch.list<int>, %arg1: !torch.int, %arg2: !torch.optional<int>) -> !torch.list<int> {\n"
67766780
" %0 = call @__torch__.torch.jit._shape_functions.unary(%arg0) : (!torch.list<int>) -> !torch.list<int>\n"
67776781
" return %0 : !torch.list<int>\n"
@@ -15367,6 +15371,18 @@ StringRef mlir::torch::Torch::getAbstractInterpLibrary() {
1536715371
" }\n"
1536815372
" return %1 : !torch.int\n"
1536915373
" }\n"
15374+
" func.func @\"__torch_mlir_dtype_fn.aten._safe_softmax\"(%arg0: !torch.tuple<int, int>, %arg1: !torch.int, %arg2: !torch.optional<int>) -> !torch.int {\n"
15375+
" %none = torch.constant.none\n"
15376+
" %0 = torch.aten.__isnot__ %arg2, %none : !torch.optional<int>, !torch.none -> !torch.bool\n"
15377+
" %1 = torch.prim.If %0 -> (!torch.int) {\n"
15378+
" %2 = torch.prim.unchecked_cast %arg2 : !torch.optional<int> -> !torch.int\n"
15379+
" torch.prim.If.yield %2 : !torch.int\n"
15380+
" } else {\n"
15381+
" %2:2 = torch.prim.TupleUnpack %arg0 : !torch.tuple<int, int> -> !torch.int, !torch.int\n"
15382+
" torch.prim.If.yield %2#1 : !torch.int\n"
15383+
" }\n"
15384+
" return %1 : !torch.int\n"
15385+
" }\n"
1537015386
" func.func @\"__torch_mlir_dtype_fn.aten._log_softmax\"(%arg0: !torch.tuple<int, int>, %arg1: !torch.int, %arg2: !torch.bool) -> !torch.int {\n"
1537115387
" %int6 = torch.constant.int 6\n"
1537215388
" %none = torch.constant.none\n"

lib/Dialect/Torch/Transforms/DecomposeComplexOps.cpp

Lines changed: 57 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -2148,6 +2148,62 @@ class DecomposeAten_SoftmaxOp : public OpRewritePattern<Aten_SoftmaxOp> {
21482148
};
21492149
} // namespace
21502150

2151+
// Ref:
2152+
// https://github.com/pytorch/pytorch/blob/5314ae2660a778b87987030182f787bb6cb092c0/aten/src/ATen/native/transformers/attention.cpp#L663-L673
2153+
namespace {
2154+
class DecomposeAten_SafeSoftmaxOp
2155+
: public OpRewritePattern<Aten_SafeSoftmaxOp> {
2156+
public:
2157+
using OpRewritePattern::OpRewritePattern;
2158+
LogicalResult matchAndRewrite(Aten_SafeSoftmaxOp op,
2159+
PatternRewriter &rewriter) const override {
2160+
BaseTensorType resultTensorType = cast<BaseTensorType>(op.getType());
2161+
if (!resultTensorType.hasDtype() || !resultTensorType.hasSizes()) {
2162+
return rewriter.notifyMatchFailure(
2163+
op, "expected result type to have sizes and dtype");
2164+
}
2165+
SmallVector<int64_t> sizes(resultTensorType.getSizes());
2166+
2167+
int64_t dimInt;
2168+
if (!matchPattern(op.getDim(), m_TorchConstantInt(&dimInt)))
2169+
return rewriter.notifyMatchFailure(op, "Unsupported: non-constant dim");
2170+
2171+
dimInt = toPositiveDim(dimInt, sizes.size());
2172+
if (!isValidDim(dimInt, sizes.size()))
2173+
return rewriter.notifyMatchFailure(op, "dim int is not valid");
2174+
2175+
Location loc = op.getLoc();
2176+
Value softmax = rewriter.create<AtenSoftmaxIntOp>(
2177+
loc, op.getType(), op.getSelf(), op.getDim(), op.getDtype());
2178+
2179+
Type resultTensorDtype = resultTensorType.getDtype();
2180+
2181+
Value negInfinity = getConstantWithGivenDtypeAndValue(
2182+
rewriter, loc, -std::numeric_limits<double>::infinity(),
2183+
resultTensorDtype);
2184+
2185+
auto boolDtype = rewriter.getI1Type();
2186+
auto boolTensorType =
2187+
resultTensorType.getWithSizesAndDtype(sizes, boolDtype);
2188+
Value masked = rewriter.create<AtenEqScalarOp>(loc, boolTensorType,
2189+
op.getSelf(), negInfinity);
2190+
2191+
sizes[dimInt] = 1;
2192+
auto maskedRowsType =
2193+
resultTensorType.getWithSizesAndDtype(sizes, boolDtype);
2194+
Value cstTrue =
2195+
rewriter.create<Torch::ConstantBoolOp>(loc, rewriter.getBoolAttr(true));
2196+
Value maskedRows = rewriter.create<AtenAllDimOp>(
2197+
loc, maskedRowsType, masked, op.getDim(), cstTrue);
2198+
Value cstZero = getConstantWithGivenDtypeAndValue(rewriter, loc, 0.0,
2199+
resultTensorDtype);
2200+
rewriter.replaceOpWithNewOp<AtenWhereScalarSelfOp>(
2201+
op, resultTensorType, maskedRows, cstZero, softmax);
2202+
return success();
2203+
}
2204+
};
2205+
} // namespace
2206+
21512207
// Aten_SoftmaxBackwardDataOp(gradOutput, output, dim) =>
21522208
// newGrad = gradOutput * output
21532209
// result = newGrad - output * sum(newGrad, dim))
@@ -9608,6 +9664,7 @@ class DecomposeComplexOpsPass
96089664
patterns);
96099665
addPatternIfTargetOpIsIllegal<DecomposeAtenSoftmaxIntOp>(patterns);
96109666
addPatternIfTargetOpIsIllegal<DecomposeAten_SoftmaxOp>(patterns);
9667+
addPatternIfTargetOpIsIllegal<DecomposeAten_SafeSoftmaxOp>(patterns);
96119668
addPatternIfTargetOpIsIllegal<DecomposeAten_LogSoftmaxOp>(patterns);
96129669
addPatternIfTargetOpIsIllegal<DecomposeAtenLogSoftmaxIntOp>(patterns);
96139670
addPatternIfTargetOpIsIllegal<DecomposeAtenLogSigmoidOp>(patterns);

lib/Dialect/Torch/Transforms/LowerToBackendContract.cpp

Lines changed: 1 addition & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -371,6 +371,7 @@ static void markDecomposedOpsAsIllegal(MLIRContext *context,
371371
llvm::StringSet<> backendLegalOpsSet) {
372372
target.addIllegalOp<AtenSoftmaxIntOp>();
373373
target.addIllegalOp<Aten_SoftmaxOp>();
374+
target.addIllegalOp<Aten_SafeSoftmaxOp>();
374375
target.addIllegalOp<Aten_LogSoftmaxOp>();
375376
target.addIllegalOp<AtenLogSoftmaxIntOp>();
376377
target.addIllegalOp<AtenLogSigmoidOp>();

projects/pt1/e2e_testing/xfail_sets.py

Lines changed: 5 additions & 8 deletions
Original file line numberDiff line numberDiff line change
@@ -504,14 +504,6 @@
504504
"ViewCollapseDynamicWithAtenSizeIntModule_basic",
505505
"ViewSizeFromOtherTensor_basic",
506506
"WeightNormInterfaceModule_basic",
507-
# REMOVE WHEN ENABLE_GQA IS ADDED
508-
"ScaledDotProductAttentionBoolMaskModule_basic",
509-
"ScaledDotProductAttentionDifferentCausalModule_basic",
510-
"ScaledDotProductAttentionDifferentModule_basic",
511-
"ScaledDotProductAttentionMaskModule_basic",
512-
"ScaledDotProductAttentionSameCausalModule_basic",
513-
"ScaledDotProductAttentionSameDynamicModule_basic",
514-
"ScaledDotProductAttentionSameModule_basic",
515507
}
516508

517509
FX_IMPORTER_CRASHING_SET = LINALG_CRASHING_SET | {
@@ -826,6 +818,9 @@
826818
"ReplicationPad2dModule_top0",
827819
"RsubInt0d_NumToTensor_Module_basic",
828820
"ScalarImplicitFloatModule_basic",
821+
# need aten.all.dim lowering to stablehlo
822+
"SafeSoftmaxModule_basic",
823+
"SafeSoftmaxNonNoneDtypeModule_basic",
829824
# REMOVE WHEN ENABLE_GQA IS ADDED
830825
"ScaledDotProductAttentionBoolMaskModule_basic",
831826
"ScaledDotProductAttentionDifferentCausalModule_basic",
@@ -2770,6 +2765,8 @@
27702765
"ReshapeAliasExpandModule_basic",
27712766
"ReshapeExpandModule_basic",
27722767
"Rot90DynamicDimsModule_basic",
2768+
"SafeSoftmaxModule_basic",
2769+
"SafeSoftmaxNonNoneDtypeModule_basic",
27732770
"ScalarConstantTupleModule_basic",
27742771
"ScalarImplicitFloatModule_basic",
27752772
"ScalarImplicitIntModule_basic",

projects/pt1/python/torch_mlir/jit_ir_importer/build_tools/abstract_interp_lib_gen.py

Lines changed: 9 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -348,6 +348,9 @@ def aten〇glu〡shape(self: List[int], dim: int = -1) -> List[int]:
348348
def aten〇_softmax〡shape(self: List[int], dim: int, half_to_float: bool) -> List[int]:
349349
return upstream_shape_functions.unary(self)
350350

351+
def aten〇_safe_softmax〡shape(self: List[int], dim: int, dtype: Optional[int] = None) -> List[int]:
352+
return upstream_shape_functions.unary(self)
353+
351354
def aten〇softmax〇int〡shape(self: List[int], dim: int, dtype: Optional[int] = None) -> List[int]:
352355
return upstream_shape_functions.unary(self)
353356

@@ -5426,6 +5429,12 @@ def aten〇_softmax〡dtype(self_rank_dtype: Tuple[int, int], dim: int, half_to_
54265429
return torch.float32
54275430
return self_dtype
54285431

5432+
def aten〇_safe_softmax〡dtype(self_rank_dtype: Tuple[int, int], dim: int, dtype: Optional[int] = None) -> int:
5433+
if dtype is not None:
5434+
return dtype
5435+
self_rank, self_dtype = self_rank_dtype
5436+
return self_dtype
5437+
54295438
@check_dtype_function(
54305439
# _check_tensors_with_the_same_dtype(num_of_tensors=1, dim=0, half_to_float=False) +
54315440
_check_tensors_with_the_same_dtype(

projects/pt1/python/torch_mlir/jit_ir_importer/build_tools/torch_ods_gen.py

Lines changed: 1 addition & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -692,6 +692,7 @@ def emit_with_mutating_variants(key, **kwargs):
692692
emit("aten::__lshift__.Scalar : (Tensor, Scalar) -> (Tensor)")
693693
emit("aten::__rshift__.Scalar : (Tensor, Scalar) -> (Tensor)")
694694
emit("aten::_softmax : (Tensor, int, bool) -> (Tensor)")
695+
emit("aten::_safe_softmax : (Tensor, int, int?) -> (Tensor)")
695696
emit("aten::mean : (Tensor, int?) -> (Tensor)")
696697
emit("aten::std : (Tensor, bool) -> (Tensor)")
697698
emit("aten::std.dim : (Tensor, int[]?, bool, bool) -> (Tensor)")

projects/pt1/python/torch_mlir_e2e_test/test_suite/basic.py

Lines changed: 46 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -1907,6 +1907,52 @@ def _LogSoftmaxModuleStable_basic(module, tu: TestUtils):
19071907
# ==============================================================================
19081908

19091909

1910+
class SafeSoftmaxModule(torch.nn.Module):
1911+
def __init__(self):
1912+
super().__init__()
1913+
1914+
@export
1915+
@annotate_args(
1916+
[
1917+
None,
1918+
([-1, -1, -1], torch.float32, True),
1919+
]
1920+
)
1921+
def forward(self, tensor):
1922+
return torch.ops.aten._safe_softmax(tensor, dim=0)
1923+
1924+
1925+
@register_test_case(module_factory=lambda: SafeSoftmaxModule())
1926+
def SafeSoftmaxModule_basic(module, tu: TestUtils):
1927+
module.forward(tu.rand(3, 2, 4))
1928+
1929+
1930+
# ==============================================================================
1931+
1932+
1933+
class SafeSoftmaxNonNoneDtypeModule(torch.nn.Module):
1934+
def __init__(self):
1935+
super().__init__()
1936+
1937+
@export
1938+
@annotate_args(
1939+
[
1940+
None,
1941+
([-1, -1, -1], torch.float32, True),
1942+
]
1943+
)
1944+
def forward(self, tensor):
1945+
return torch.ops.aten._safe_softmax(tensor, dim=2, dtype=torch.float64)
1946+
1947+
1948+
@register_test_case(module_factory=lambda: SafeSoftmaxNonNoneDtypeModule())
1949+
def SafeSoftmaxNonNoneDtypeModule_basic(module, tu: TestUtils):
1950+
module.forward(tu.rand(3, 2, 4))
1951+
1952+
1953+
# ==============================================================================
1954+
1955+
19101956
class SoftplusModule(torch.nn.Module):
19111957
def __init__(self):
19121958
super().__init__()

0 commit comments

Comments
 (0)