// Copyright (c) 2021 PaddlePaddle Authors. All Rights Reserved. // // Licensed under the Apache License, Version 2.0 (the "License"); // you may not use this file except in compliance with the License. // You may obtain a copy of the License at // // http://www.apache.org/licenses/LICENSE-2.0 // // Unless required by applicable law or agreed to in writing, software // distributed under the License is distributed on an "AS IS" BASIS, // WITHOUT WARRANTIES OR CONDITIONS OF ANY KIND, either express or implied. // See the License for the specific language governing permissions and // limitations under the License. #include "paddle/fluid/framework/infershape_utils.h" #include "paddle/fluid/framework/op_registry.h" #include "paddle/phi/core/infermeta_utils.h" #include "paddle/phi/infermeta/unary.h" namespace paddle { namespace operators { class DiagonalOp : public framework::OperatorWithKernel { public: using framework::OperatorWithKernel::OperatorWithKernel; }; class DiagonalOpMaker : public framework::OpProtoAndCheckerMaker { public: void Make() override { AddInput("Input", "(Tensor) The input tensor, from which the diagonals are taken."); AddOutput( "Out", "(Tensor) The partial view of input with the its diagonal elements."); AddAttr( "offset", R"DOC((int, default 0), offset of the diagonal from the main diagonal. Can be both positive and negative. Default: 0. )DOC") .SetDefault(0); AddAttr( "axis1", R"DOC((int, default 0), the first axis of the 2-D planes from which the diagonals should be taken. Can be either positive or negative. Default: 0. )DOC") .SetDefault(0); AddAttr( "axis2", R"DOC((int, default 1), the second axis of the 2-D planes from which the diagonals should be taken. Can be either positive or negative. Default: 1. )DOC") .SetDefault(1); AddComment(R"DOC( Diagonal Operator. Return a partial view of input with the its diagonal elements of the input tensor. The behavior of this operator is similar to how `numpy.diagonal` works. )DOC"); } }; class DiagonalGradOp : public framework::OperatorWithKernel { public: using framework::OperatorWithKernel::OperatorWithKernel; void InferShape(framework::InferShapeContext *ctx) const override { OP_INOUT_CHECK(ctx->HasInput("Input"), "Input", "Input", "DiagonalGrad"); OP_INOUT_CHECK(ctx->HasOutput(framework::GradVarName("Input")), "Output", framework::GradVarName("Input"), "DiagonalGrad"); ctx->SetOutputDim(framework::GradVarName("Input"), ctx->GetInputDim("Input")); } protected: framework::OpKernelType GetExpectedKernelType( const framework::ExecutionContext &ctx) const override { return framework::OpKernelType(OperatorWithKernel::IndicateVarDataType( ctx, framework::GradVarName("Out")), ctx.GetPlace()); } }; template class DiagonalGradOpMaker : public framework::SingleGradOpMaker { public: using framework::SingleGradOpMaker::SingleGradOpMaker; protected: void Apply(GradOpPtr grad_op) const override { grad_op->SetType("diagonal_grad"); grad_op->SetInput("Input", this->Input("Input")); grad_op->SetInput(framework::GradVarName("Out"), this->OutputGrad("Out")); grad_op->SetOutput(framework::GradVarName("Input"), this->InputGrad("Input")); grad_op->SetAttrMap(this->Attrs()); } }; DECLARE_NO_NEED_BUFFER_VARS_INFERER(DiagonalGradNoNeedBufferVarsInferer, "Input"); } // namespace operators } // namespace paddle namespace ops = paddle::operators; DELCARE_INFER_SHAPE_FUNCTOR(diagonal, DiagonalInferShapeFunctor, PT_INFER_META(phi::DiagonalInferMeta)); REGISTER_OPERATOR(diagonal, ops::DiagonalOp, ops::DiagonalOpMaker, ops::DiagonalGradOpMaker, ops::DiagonalGradOpMaker, DiagonalInferShapeFunctor); REGISTER_OPERATOR(diagonal_grad, ops::DiagonalGradOp, ops::DiagonalGradNoNeedBufferVarsInferer)