Skip to content

Commit 8ae4dbf

Browse files
author
Prashant Kumar
committed
Add aten::nll_loss_backward op
The lowering of aten::nll_loss_backward op has been added from torch to linalg dialect. The changes has been made as a part of -torch-convert-to-linalg pass. Signed-off-by: Prashant Kumar prashant@nod-labs.com
1 parent 0079901 commit 8ae4dbf

File tree

5 files changed

+216
-2
lines changed

5 files changed

+216
-2
lines changed

e2e_testing/torchscript/nll_loss.py

Lines changed: 58 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -60,3 +60,61 @@ def forward(self, x, y):
6060
@register_test_case(module_factory=lambda: NllLossModule_ignore_index_out_of_bounds())
6161
def NllLossModule_ignore_index(module, tu: TestUtils):
6262
module.forward(tu.rand(2, 3), torch.tensor([0, 1]))
63+
64+
class NllLossModule_backward(torch.nn.Module):
65+
66+
def __init__(self):
67+
super().__init__()
68+
69+
@export
70+
@annotate_args([
71+
None,
72+
([-1], torch.float32, True),
73+
([-1, -1], torch.float32, True),
74+
([-1], torch.int64, True),
75+
([], torch.float32, True),
76+
])
77+
def forward(self, grad_output, input, target, total_weight):
78+
return torch.ops.aten.nll_loss_backward(grad_output=grad_output,
79+
self=input,
80+
target=target,
81+
weight=None,
82+
reduction=0,
83+
ignore_index=10,
84+
total_weight=total_weight)
85+
86+
87+
@register_test_case(module_factory=lambda: NllLossModule_backward())
88+
def NllLossModuleBackward_basic(module, tu: TestUtils):
89+
module.forward(tu.rand(3), tu.rand(3, 4), torch.tensor([2, 3, 0]),
90+
torch.tensor(3.))
91+
92+
93+
class NllLossModule_backward_ignore_index(torch.nn.Module):
94+
95+
def __init__(self):
96+
super().__init__()
97+
98+
@export
99+
@annotate_args([
100+
None,
101+
([-1], torch.float32, True),
102+
([-1, -1], torch.float32, True),
103+
([-1], torch.int64, True),
104+
([], torch.float32, True),
105+
])
106+
def forward(self, grad_output, input, target, total_weight):
107+
return torch.ops.aten.nll_loss_backward(grad_output=grad_output,
108+
self=input,
109+
target=target,
110+
weight=None,
111+
reduction=0,
112+
ignore_index=1,
113+
total_weight=total_weight)
114+
115+
116+
@register_test_case(
117+
module_factory=lambda: NllLossModule_backward_ignore_index())
118+
def NllLossModuleBackward_ignore_index(module, tu: TestUtils):
119+
module.forward(tu.rand(3), tu.rand(3, 4), torch.tensor([2, 3, 0]),
120+
torch.tensor(3.))

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

Lines changed: 20 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -1852,6 +1852,26 @@ def Torch_AtenNllLossForwardOp : Torch_Op<"aten.nll_loss_forward", [
18521852
let assemblyFormat = "$self `,` $target `,` $weight `,` $reduction `,` $ignore_index attr-dict `:` qualified(type($self)) `,` qualified(type($target)) `,` qualified(type($weight)) `,` qualified(type($reduction)) `,` qualified(type($ignore_index)) `->` qualified(type($output)) `,` qualified(type($total_weight))";
18531853
}
18541854

1855+
def Torch_AtenNllLossBackwardOp : Torch_Op<"aten.nll_loss_backward", [
1856+
AllowsTypeRefinement,
1857+
HasValueSemantics
1858+
]> {
1859+
let summary = "Generated op for `aten::nll_loss_backward : (Tensor, Tensor, Tensor, Tensor?, int, int, Tensor) -> (Tensor)`";
1860+
let arguments = (ins
1861+
AnyTorchTensorType:$grad_output,
1862+
AnyTorchTensorType:$self,
1863+
AnyTorchTensorType:$target,
1864+
AnyTorchOptionalTensorType:$weight,
1865+
Torch_IntType:$reduction,
1866+
Torch_IntType:$ignore_index,
1867+
AnyTorchTensorType:$total_weight
1868+
);
1869+
let results = (outs
1870+
AnyTorchTensorType:$result
1871+
);
1872+
let assemblyFormat = "$grad_output `,` $self `,` $target `,` $weight `,` $reduction `,` $ignore_index `,` $total_weight attr-dict `:` qualified(type($grad_output)) `,` qualified(type($self)) `,` qualified(type($target)) `,` qualified(type($weight)) `,` qualified(type($reduction)) `,` qualified(type($ignore_index)) `,` qualified(type($total_weight)) `->` qualified(type($result))";
1873+
}
1874+
18551875
def Torch_AtenConstantPadNdOp : Torch_Op<"aten.constant_pad_nd", [
18561876
AllowsTypeRefinement,
18571877
HasValueSemantics

lib/Conversion/TorchToLinalg/TorchToLinalg.cpp

Lines changed: 109 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -328,6 +328,11 @@ static Value buildUnitNormalCdf(OpBuilder &b, Location &loc, Value x) {
328328
return buildNormalCdf(b, loc, x, zero, one);
329329
}
330330

331+
// These constants control the reduction behavior of the loss functions.
332+
// None, Mean and Sum corresponds to "do not reduce", "Mean of losses", and "sum
333+
// of losses" respectively.
334+
enum Reduction { None, Mean, Sum, END };
335+
331336
namespace {
332337
class ConvertAtenAdaptiveAvgPool2dOp
333338
: public OpConversionPattern<AtenAdaptiveAvgPool2dOp> {
@@ -1323,6 +1328,108 @@ class ConvertAtenNllLossForwardOp
13231328
};
13241329
} // namespace
13251330

1331+
// Given `grad_output`, `input`, `target`, `nll_loss_backward` is given by:
1332+
// for i in range(0, len(input[0])):
1333+
// for j in range(0, len(input[1])):
1334+
// nll_loss_backward[i][j] = (j == target[i]) ? -grad_output[i] : 0
1335+
// TODO: `weight` and `reduction` operands are still to be taken care of.
1336+
namespace {
1337+
class ConvertAtenNllLossBackwardOp
1338+
: public OpConversionPattern<AtenNllLossBackwardOp> {
1339+
public:
1340+
using OpConversionPattern::OpConversionPattern;
1341+
LogicalResult
1342+
matchAndRewrite(AtenNllLossBackwardOp op, OpAdaptor adaptor,
1343+
ConversionPatternRewriter &rewriter) const override {
1344+
if (failed(verifyLinalgCompatibleTypes(op, rewriter)))
1345+
return failure();
1346+
Location loc = op->getLoc();
1347+
Value input = adaptor.self();
1348+
Value target = adaptor.target();
1349+
Value weight = adaptor.weight();
1350+
Value grad_output = adaptor.grad_output();
1351+
1352+
int64_t reduction;
1353+
if (!matchPattern(op.reduction(), m_TorchConstantInt(&reduction)))
1354+
return rewriter.notifyMatchFailure(op, "dim must be constant");
1355+
1356+
// TODO: Handle reduction.
1357+
if (reduction != Reduction::None)
1358+
return rewriter.notifyMatchFailure(
1359+
op, "reduction along dimensions is not supported.");
1360+
1361+
// TODO: Incorporate the weight argument.
1362+
if (!weight.getType().isa<Torch::NoneType>())
1363+
return rewriter.notifyMatchFailure(
1364+
op, "Unimplemented, the weight operand is not incorporated.");
1365+
1366+
Value ignoreIndex = adaptor.ignore_index();
1367+
Value ignoreIndexVal = castIntToIndex(rewriter, loc, ignoreIndex);
1368+
1369+
unsigned inputRank = input.getType().cast<RankedTensorType>().getRank();
1370+
unsigned targetRank = target.getType().cast<RankedTensorType>().getRank();
1371+
1372+
// TODO: Cases with targetRank != 1 where `Mean` or `Sum` reduction is
1373+
// required.
1374+
if (inputRank != 2 || targetRank != 1) {
1375+
return rewriter.notifyMatchFailure(
1376+
op, "expected input and target to be rank 2 and 1 respectively");
1377+
}
1378+
RankedTensorType resultType = getTypeConverter()
1379+
->convertType(op->getResult(0).getType())
1380+
.cast<RankedTensorType>();
1381+
1382+
Type elementType = resultType.getElementType();
1383+
1384+
// Given there is no reduction `grad_input` size is equal to `input` size.
1385+
auto outputSize = getTensorSizes(rewriter, loc, input);
1386+
Value initTensor0 =
1387+
createZeroInitTensor(rewriter, loc, outputSize, elementType);
1388+
Value zeroVal = rewriter.create<arith::ConstantOp>(
1389+
loc, rewriter.getZeroAttr(elementType));
1390+
1391+
SmallVector<AffineExpr> targetExpr{rewriter.getAffineDimExpr(0)};
1392+
SmallVector<AffineExpr> resultExpr{rewriter.getAffineDimExpr(0),
1393+
rewriter.getAffineDimExpr(1)};
1394+
SmallVector<StringRef> iteratorTypes{getParallelIteratorTypeName(),
1395+
getParallelIteratorTypeName()};
1396+
auto indexingMaps =
1397+
AffineMap::inferFromExprList({targetExpr, targetExpr, resultExpr});
1398+
Value finalRes =
1399+
rewriter
1400+
.create<linalg::GenericOp>(
1401+
loc, resultType, ValueRange{target, grad_output}, initTensor0,
1402+
/*indexingMaps=*/indexingMaps,
1403+
/*iteratorTypes=*/iteratorTypes,
1404+
[&](OpBuilder &b, Location loc, ValueRange args) {
1405+
Value indTarget = rewriter.create<arith::IndexCastOp>(
1406+
loc, rewriter.getIndexType(), args[0]);
1407+
Value indJ = rewriter.create<linalg::IndexOp>(loc, 1);
1408+
1409+
// The final result is given by:
1410+
// grad_input[i][j] = (j == target[i]) ? -grad_output[i] : 0
1411+
Value cmpEq = rewriter.create<arith::CmpIOp>(
1412+
loc, arith::CmpIPredicate::eq, indJ, indTarget);
1413+
1414+
// The target index shouldn't be equal to `ignoreIndex`.
1415+
Value cmpNEq = rewriter.create<arith::CmpIOp>(
1416+
loc, arith::CmpIPredicate::ne, ignoreIndexVal, indTarget);
1417+
Value finalPredicate =
1418+
rewriter.create<arith::AndIOp>(loc, cmpEq, cmpNEq);
1419+
Value negate =
1420+
rewriter.create<arith::NegFOp>(loc, elementType, args[1]);
1421+
Value selectFinal = rewriter.create<mlir::SelectOp>(
1422+
loc, finalPredicate, negate, zeroVal);
1423+
b.create<linalg::YieldOp>(loc, selectFinal);
1424+
})
1425+
.getResult(0);
1426+
1427+
rewriter.replaceOp(op, finalRes);
1428+
return success();
1429+
}
1430+
};
1431+
} // namespace
1432+
13261433
namespace {
13271434
// See comments at in convertMmOp and the heading for this section for general
13281435
// considerations. This function needs to be auto-generated.
@@ -4525,6 +4632,8 @@ class ConvertTorchToLinalg
45254632
patterns.add<ConvertAtenSliceTensorOp>(typeConverter, context);
45264633
target.addIllegalOp<AtenNllLossForwardOp>();
45274634
patterns.add<ConvertAtenNllLossForwardOp>(typeConverter, context);
4635+
target.addIllegalOp<AtenNllLossBackwardOp>();
4636+
patterns.add<ConvertAtenNllLossBackwardOp>(typeConverter, context);
45284637
target.addIllegalOp<AtenIndexSelectOp>();
45294638
patterns.add<ConvertAtenIndexSelectOp>(typeConverter, context);
45304639
patterns.add<ConvertAtenScalarToTensorLike>(typeConverter, context);

lib/Dialect/Torch/Transforms/RefineTypes.cpp

Lines changed: 28 additions & 2 deletions
Original file line numberDiff line numberDiff line change
@@ -46,6 +46,11 @@ static ScalarType getScalarTypeForType(Type type) {
4646
llvm::report_fatal_error("unhandled type for getScalarTypeForType");
4747
}
4848

49+
// These constants control the reduction behavior of the loss functions.
50+
// None, Mean and Sum corresponds to "do not reduce", "Mean of losses", and "sum
51+
// of losses" respectively.
52+
enum Reduction { None, Mean, Sum, END };
53+
4954
static Type getTypeForScalarType(MLIRContext *context, ScalarType dtypeInt) {
5055
switch (dtypeInt) {
5156
case ScalarType::Float:
@@ -489,6 +494,8 @@ class TypeAnalyzer : public ForwardDataFlowAnalysis<ValueKnowledge> {
489494
return visitBinaryScalarOp(op, operands);
490495
} else if (auto nllForwardOp = dyn_cast<AtenNllLossForwardOp>(op)) {
491496
return visitAtenNllLossForwardOp(nllForwardOp, operands);
497+
} else if (auto nllBackwardOp = dyn_cast<AtenNllLossBackwardOp>(op)) {
498+
return visitAtenNllLossBackwardOp(nllBackwardOp, operands);
492499
} else if (auto nativeLayerNormOp = dyn_cast<AtenNativeLayerNormOp>(op)) {
493500
return visitAtenNativeLayerNormOp(nativeLayerNormOp, operands);
494501
} else if (auto constantPadNdOp = dyn_cast<AtenConstantPadNdOp>(op)) {
@@ -647,6 +654,9 @@ class TypeAnalyzer : public ForwardDataFlowAnalysis<ValueKnowledge> {
647654
ChangeResult visitAtenNllLossForwardOp(
648655
AtenNllLossForwardOp op,
649656
ArrayRef<LatticeElement<ValueKnowledge> *> operands);
657+
ChangeResult visitAtenNllLossBackwardOp(
658+
AtenNllLossBackwardOp op,
659+
ArrayRef<LatticeElement<ValueKnowledge> *> operands);
650660
ChangeResult visitAtenNativeLayerNormOp(
651661
AtenNativeLayerNormOp op,
652662
ArrayRef<LatticeElement<ValueKnowledge> *> operands);
@@ -1188,8 +1198,8 @@ ChangeResult TypeAnalyzer::visitAtenNllLossForwardOp(
11881198

11891199
if (self.hasSizes &&
11901200
matchPattern(op.reduction(), m_TorchConstantInt(&reduction))) {
1191-
// reduction == 1 means reduce 1st dim.
1192-
resultRank = reduction == 1 ? resultRank - 1 : resultRank;
1201+
if (reduction != Reduction::None)
1202+
resultRank -= 1;
11931203
}
11941204
outputKnowledge.sizes.resize(resultRank - 1, kUnknownSize);
11951205
outputKnowledge.hasSizes = true;
@@ -1199,6 +1209,22 @@ ChangeResult TypeAnalyzer::visitAtenNllLossForwardOp(
11991209
return resultLattice;
12001210
}
12011211

1212+
ChangeResult TypeAnalyzer::visitAtenNllLossBackwardOp(
1213+
AtenNllLossBackwardOp op,
1214+
ArrayRef<LatticeElement<ValueKnowledge> *> operands) {
1215+
auto self = operands[1]->getValue();
1216+
auto knowledge =
1217+
ValueKnowledge::getNotNonePessimisticValueState(op.getContext());
1218+
1219+
knowledge.dtype = self.dtype;
1220+
if (self.hasSizes) {
1221+
unsigned resultRank = self.sizes.size();
1222+
knowledge.sizes.resize(resultRank, kUnknownSize);
1223+
knowledge.hasSizes = true;
1224+
}
1225+
return getLatticeElement(op.getResult()).join(knowledge);
1226+
}
1227+
12021228
ChangeResult TypeAnalyzer::visitAtenSqueezeDimOp(
12031229
AtenSqueezeDimOp op, ArrayRef<LatticeElement<ValueKnowledge> *> operands) {
12041230
auto operand = operands[0]->getValue();

python/torch_mlir/dialects/torch/importer/jit_ir/build_tools/torch_ods_gen.py

Lines changed: 1 addition & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -548,6 +548,7 @@ def emit_with_mutating_variants(key, **kwargs):
548548
emit("aten::std : (Tensor, bool) -> (Tensor)")
549549
emit("aten::var : (Tensor, bool) -> (Tensor)")
550550
emit("aten::nll_loss_forward : (Tensor, Tensor, Tensor?, int, int) -> (Tensor, Tensor)")
551+
emit("aten::nll_loss_backward : (Tensor, Tensor, Tensor, Tensor?, int, int, Tensor) -> (Tensor)")
551552

552553
# Misc tensor ops.
553554
emit("aten::constant_pad_nd : (Tensor, int[], Scalar) -> (Tensor)")

0 commit comments

Comments
 (0)